loss.py 9.8 KB

123456789101112131415161718192021222324252627282930313233343536373839404142434445464748495051525354555657585960616263646566676869707172737475767778798081828384858687888990919293949596979899100101102103104105106107108109110111112113114115116117118119120121122123124125126127128129130131132133134135136137138139140141142143144145146147148149150151152153154155156157158159160161162163164165166167168169170171172173174175176177178179180181182183184185186187188189190191192193194195196197198199200201202203204205206207208209210211212213214
  1. import copy
  2. import torch
  3. import torch.nn as nn
  4. import torch.nn.functional as F
  5. try:
  6. from .loss_utils import sigmoid_focal_loss
  7. from .loss_utils import box_cxcywh_to_xyxy, box_xyxy_to_cxcywh, generalized_box_iou, bbox2delta
  8. from .loss_utils import is_dist_avail_and_initialized, get_world_size
  9. from .matcher import HungarianMatcher
  10. except:
  11. from loss_utils import sigmoid_focal_loss
  12. from loss_utils import box_cxcywh_to_xyxy, box_xyxy_to_cxcywh, generalized_box_iou, bbox2delta
  13. from loss_utils import is_dist_avail_and_initialized, get_world_size
  14. from matcher import HungarianMatcher
  15. class Criterion(nn.Module):
  16. """ This class computes the loss for DETR.
  17. The process happens in two steps:
  18. 1) we compute hungarian assignment between ground truth boxes and the outputs of the model
  19. 2) we supervise each pair of matched ground-truth / prediction (supervise class and box)
  20. """
  21. def __init__(self, cfg, num_classes=80, aux_loss=False):
  22. super().__init__()
  23. # ------------ Basic parameters ------------
  24. self.cfg = cfg
  25. self.num_classes = num_classes
  26. self.k_one2many = cfg['k_one2many']
  27. self.lambda_one2many = cfg['lambda_one2many']
  28. self.aux_loss = aux_loss
  29. self.losses = ['labels', 'boxes']
  30. # ------------- Focal loss -------------
  31. self.alpha = 0.25
  32. self.gamma = 2.0
  33. # ------------ Matcher ------------
  34. self.matcher = HungarianMatcher(cost_class = cfg['matcher_hpy']['cost_class'],
  35. cost_bbox = cfg['matcher_hpy']['cost_bbox'],
  36. cost_giou = cfg['matcher_hpy']['cost_giou']
  37. )
  38. # ------------- Loss weight -------------
  39. self.weight_dict = {'loss_cls': cfg['loss_coeff']['class'],
  40. 'loss_box': cfg['loss_coeff']['bbox'],
  41. 'loss_giou': cfg['loss_coeff']['giou']}
  42. def _get_src_permutation_idx(self, indices):
  43. # permute predictions following indices
  44. batch_idx = torch.cat([torch.full_like(src, i) for i, (src, _) in enumerate(indices)])
  45. src_idx = torch.cat([src for (src, _) in indices])
  46. return batch_idx, src_idx
  47. def _get_tgt_permutation_idx(self, indices):
  48. # permute targets following indices
  49. batch_idx = torch.cat([torch.full_like(tgt, i) for i, (_, tgt) in enumerate(indices)])
  50. tgt_idx = torch.cat([tgt for (_, tgt) in indices])
  51. return batch_idx, tgt_idx
  52. def loss_labels(self, outputs, targets, indices, num_boxes):
  53. """Classification loss (NLL)
  54. targets dicts must contain the key "labels" containing a tensor of dim [nb_target_boxes]
  55. """
  56. assert 'pred_logits' in outputs
  57. src_logits = outputs['pred_logits']
  58. # prepare class targets
  59. idx = self._get_src_permutation_idx(indices)
  60. target_classes_o = torch.cat([t["labels"][J] for t, (_, J) in zip(targets, indices)]).to(src_logits.device)
  61. target_classes = torch.full(src_logits.shape[:2],
  62. self.num_classes,
  63. dtype=torch.int64,
  64. device=src_logits.device)
  65. target_classes[idx] = target_classes_o
  66. # to one-hot labels
  67. target_classes_onehot = torch.zeros([*src_logits.shape[:2], self.num_classes + 1],
  68. dtype=src_logits.dtype,
  69. layout=src_logits.layout,
  70. device=src_logits.device)
  71. target_classes_onehot.scatter_(2, target_classes.unsqueeze(-1), 1)
  72. target_classes_onehot = target_classes_onehot[..., :-1]
  73. # focal loss
  74. loss_cls = sigmoid_focal_loss(src_logits, target_classes_onehot, num_boxes, self.alpha, self.gamma)
  75. losses = {}
  76. losses['loss_cls'] = loss_cls * src_logits.shape[1]
  77. return losses
  78. def loss_boxes(self, outputs, targets, indices, num_boxes):
  79. """Compute the losses related to the bounding boxes, the L1 regression loss and the GIoU loss
  80. targets dicts must contain the key "boxes" containing a tensor of dim [nb_target_boxes, 4]
  81. The target boxes are expected in format (center_x, center_y, w, h), normalized by the image size.
  82. """
  83. assert 'pred_boxes' in outputs
  84. # prepare bbox targets
  85. idx = self._get_src_permutation_idx(indices)
  86. src_boxes = outputs['pred_boxes'][idx]
  87. target_boxes = torch.cat([t['boxes'][i] for t, (_, i) in zip(targets, indices)], dim=0).to(src_boxes.device)
  88. # compute L1 loss
  89. loss_bbox = F.l1_loss(src_boxes, box_xyxy_to_cxcywh(target_boxes), reduction='none')
  90. src_deltas = outputs["pred_deltas"][idx]
  91. src_boxes_old = outputs["pred_boxes_old"][idx]
  92. target_deltas = bbox2delta(src_boxes_old, target_boxes)
  93. loss_bbox = F.l1_loss(src_deltas, target_deltas, reduction="none")
  94. # compute GIoU loss
  95. bbox_giou = generalized_box_iou(box_cxcywh_to_xyxy(src_boxes),
  96. box_cxcywh_to_xyxy(target_boxes))
  97. loss_giou = 1 - torch.diag(bbox_giou)
  98. losses = {}
  99. losses['loss_box'] = loss_bbox.sum() / num_boxes
  100. losses['loss_giou'] = loss_giou.sum() / num_boxes
  101. return losses
  102. def get_loss(self, loss, outputs, targets, indices, num_boxes, **kwargs):
  103. loss_map = {
  104. 'labels': self.loss_labels,
  105. 'boxes': self.loss_boxes,
  106. }
  107. assert loss in loss_map, f'do you really want to compute {loss} loss?'
  108. return loss_map[loss](outputs, targets, indices, num_boxes, **kwargs)
  109. def compute_loss(self, outputs, targets):
  110. """ This performs the loss computation.
  111. Parameters:
  112. outputs: dict of tensors, see the output specification of the model for the format
  113. targets: list of dicts, such that len(targets) == batch_size.
  114. The expected keys in each dict depends on the losses applied, see each loss' doc
  115. """
  116. outputs_without_aux = {
  117. k: v
  118. for k, v in outputs.items()
  119. if k != "aux_outputs" and k != "enc_outputs"
  120. }
  121. # Retrieve the matching between the outputs of the last layer and the targets
  122. indices = self.matcher(outputs_without_aux, targets)
  123. # Compute the average number of target boxes accross all nodes, for normalization purposes
  124. num_boxes = sum(len(t["labels"]) for t in targets)
  125. num_boxes = torch.as_tensor(
  126. [num_boxes], dtype=torch.float, device=next(iter(outputs.values())).device
  127. )
  128. if is_dist_avail_and_initialized():
  129. torch.distributed.all_reduce(num_boxes)
  130. num_boxes = torch.clamp(num_boxes / get_world_size(), min=1).item()
  131. # Compute all the requested losses
  132. losses = {}
  133. for loss in self.losses:
  134. kwargs = {}
  135. l_dict = self.get_loss(loss, outputs, targets, indices, num_boxes, **kwargs)
  136. l_dict = {k: l_dict[k] * self.weight_dict[k] for k in l_dict if k in self.weight_dict}
  137. losses.update(l_dict)
  138. # In case of auxiliary losses, we repeat this process with the output of each intermediate layer.
  139. if "aux_outputs" in outputs:
  140. for i, aux_outputs in enumerate(outputs["aux_outputs"]):
  141. indices = self.matcher(aux_outputs, targets)
  142. for loss in self.losses:
  143. kwargs = {}
  144. l_dict = self.get_loss(loss, aux_outputs, targets, indices, num_boxes, **kwargs)
  145. l_dict = {k: l_dict[k] * self.weight_dict[k] for k in l_dict if k in self.weight_dict}
  146. l_dict = {k + f"_{i}": v for k, v in l_dict.items()}
  147. losses.update(l_dict)
  148. if "enc_outputs" in outputs:
  149. enc_outputs = outputs["enc_outputs"]
  150. bin_targets = copy.deepcopy(targets)
  151. for bt in bin_targets:
  152. bt["labels"] = torch.zeros_like(bt["labels"])
  153. indices = self.matcher(enc_outputs, bin_targets)
  154. for loss in self.losses:
  155. kwargs = {}
  156. l_dict = self.get_loss(loss, enc_outputs, bin_targets, indices, num_boxes, **kwargs)
  157. l_dict = {k: l_dict[k] * self.weight_dict[k] for k in l_dict if k in self.weight_dict}
  158. l_dict = {k + "_enc": v for k, v in l_dict.items()}
  159. losses.update(l_dict)
  160. return losses
  161. def forward(self, outputs, targets):
  162. # --------------------- One-to-one losses ---------------------
  163. outputs_one2one = {k: v for k, v in outputs.items() if "one2many" not in k}
  164. loss_dict = self.compute_loss(outputs_one2one, targets)
  165. # --------------------- One-to-many losses ---------------------
  166. outputs_one2many = {k[:-9]: v for k, v in outputs.items() if "one2many" in k}
  167. if len(outputs_one2many) > 0:
  168. # Copy targets
  169. multi_targets = copy.deepcopy(targets)
  170. for target in multi_targets:
  171. target["boxes"] = target["boxes"].repeat(self.k_one2many, 1)
  172. target["labels"] = target["labels"].repeat(self.k_one2many)
  173. # Compute one-to-many losses
  174. one2many_loss_dict = self.compute_loss(outputs_one2many, multi_targets)
  175. # add one2many losses in to the final loss_dict
  176. for k, v in one2many_loss_dict.items():
  177. if k + "_one2many" in loss_dict.keys():
  178. loss_dict[k + "_one2many"] += v * self.lambda_one2many
  179. else:
  180. loss_dict[k + "_one2many"] = v * self.lambda_one2many
  181. return loss_dict
  182. # build criterion
  183. def build_criterion(cfg, num_classes, aux_loss=True):
  184. criterion = Criterion(cfg, num_classes, aux_loss)
  185. return criterion