varifocal_loss.py 5.6 KB

123456789101112131415161718192021222324252627282930313233343536373839404142434445464748495051525354555657585960616263646566676869707172737475767778798081828384858687888990919293949596979899100101102103104105106107108109110111112113114115116117118119120121122123124125126127128129130131132133134135136137138139140141
  1. # Copyright (c) OpenMMLab. All rights reserved.
  2. from typing import Optional
  3. import torch.nn as nn
  4. import torch.nn.functional as F
  5. from torch import Tensor
  6. from mmdet.registry import MODELS
  7. from .utils import weight_reduce_loss
  8. def varifocal_loss(pred: Tensor,
  9. target: Tensor,
  10. weight: Optional[Tensor] = None,
  11. alpha: float = 0.75,
  12. gamma: float = 2.0,
  13. iou_weighted: bool = True,
  14. reduction: str = 'mean',
  15. avg_factor: Optional[int] = None) -> Tensor:
  16. """`Varifocal Loss <https://arxiv.org/abs/2008.13367>`_
  17. Args:
  18. pred (Tensor): The prediction with shape (N, C), C is the
  19. number of classes.
  20. target (Tensor): The learning target of the iou-aware
  21. classification score with shape (N, C), C is the number of classes.
  22. weight (Tensor, optional): The weight of loss for each
  23. prediction. Defaults to None.
  24. alpha (float, optional): A balance factor for the negative part of
  25. Varifocal Loss, which is different from the alpha of Focal Loss.
  26. Defaults to 0.75.
  27. gamma (float, optional): The gamma for calculating the modulating
  28. factor. Defaults to 2.0.
  29. iou_weighted (bool, optional): Whether to weight the loss of the
  30. positive example with the iou target. Defaults to True.
  31. reduction (str, optional): The method used to reduce the loss into
  32. a scalar. Defaults to 'mean'. Options are "none", "mean" and
  33. "sum".
  34. avg_factor (int, optional): Average factor that is used to average
  35. the loss. Defaults to None.
  36. Returns:
  37. Tensor: Loss tensor.
  38. """
  39. # pred and target should be of the same size
  40. assert pred.size() == target.size()
  41. pred_sigmoid = pred.sigmoid()
  42. target = target.type_as(pred)
  43. if iou_weighted:
  44. focal_weight = target * (target > 0.0).float() + \
  45. alpha * (pred_sigmoid - target).abs().pow(gamma) * \
  46. (target <= 0.0).float()
  47. else:
  48. focal_weight = (target > 0.0).float() + \
  49. alpha * (pred_sigmoid - target).abs().pow(gamma) * \
  50. (target <= 0.0).float()
  51. loss = F.binary_cross_entropy_with_logits(
  52. pred, target, reduction='none') * focal_weight
  53. loss = weight_reduce_loss(loss, weight, reduction, avg_factor)
  54. return loss
  55. @MODELS.register_module()
  56. class VarifocalLoss(nn.Module):
  57. def __init__(self,
  58. use_sigmoid: bool = True,
  59. alpha: float = 0.75,
  60. gamma: float = 2.0,
  61. iou_weighted: bool = True,
  62. reduction: str = 'mean',
  63. loss_weight: float = 1.0) -> None:
  64. """`Varifocal Loss <https://arxiv.org/abs/2008.13367>`_
  65. Args:
  66. use_sigmoid (bool, optional): Whether the prediction is
  67. used for sigmoid or softmax. Defaults to True.
  68. alpha (float, optional): A balance factor for the negative part of
  69. Varifocal Loss, which is different from the alpha of Focal
  70. Loss. Defaults to 0.75.
  71. gamma (float, optional): The gamma for calculating the modulating
  72. factor. Defaults to 2.0.
  73. iou_weighted (bool, optional): Whether to weight the loss of the
  74. positive examples with the iou target. Defaults to True.
  75. reduction (str, optional): The method used to reduce the loss into
  76. a scalar. Defaults to 'mean'. Options are "none", "mean" and
  77. "sum".
  78. loss_weight (float, optional): Weight of loss. Defaults to 1.0.
  79. """
  80. super().__init__()
  81. assert use_sigmoid is True, \
  82. 'Only sigmoid varifocal loss supported now.'
  83. assert alpha >= 0.0
  84. self.use_sigmoid = use_sigmoid
  85. self.alpha = alpha
  86. self.gamma = gamma
  87. self.iou_weighted = iou_weighted
  88. self.reduction = reduction
  89. self.loss_weight = loss_weight
  90. def forward(self,
  91. pred: Tensor,
  92. target: Tensor,
  93. weight: Optional[Tensor] = None,
  94. avg_factor: Optional[int] = None,
  95. reduction_override: Optional[str] = None) -> Tensor:
  96. """Forward function.
  97. Args:
  98. pred (Tensor): The prediction with shape (N, C), C is the
  99. number of classes.
  100. target (Tensor): The learning target of the iou-aware
  101. classification score with shape (N, C), C is
  102. the number of classes.
  103. weight (Tensor, optional): The weight of loss for each
  104. prediction. Defaults to None.
  105. avg_factor (int, optional): Average factor that is used to average
  106. the loss. Defaults to None.
  107. reduction_override (str, optional): The reduction method used to
  108. override the original reduction method of the loss.
  109. Options are "none", "mean" and "sum".
  110. Returns:
  111. Tensor: The calculated loss
  112. """
  113. assert reduction_override in (None, 'none', 'mean', 'sum')
  114. reduction = (
  115. reduction_override if reduction_override else self.reduction)
  116. if self.use_sigmoid:
  117. loss_cls = self.loss_weight * varifocal_loss(
  118. pred,
  119. target,
  120. weight,
  121. alpha=self.alpha,
  122. gamma=self.gamma,
  123. iou_weighted=self.iou_weighted,
  124. reduction=reduction,
  125. avg_factor=avg_factor)
  126. else:
  127. raise NotImplementedError
  128. return loss_cls