loss_d_fine.py 15 KB

123456789101112131415161718192021222324252627282930313233343536373839404142434445464748495051525354555657585960616263646566676869707172737475767778798081828384858687888990919293949596979899100101102103104105106107108109110111112113114115116117118119120121122123124125126127128129130131132133134135136137138139140141142143144145146147148149150151152153154155156157158159160161162163164165166167168169170171172173174175176177178179180181182183184185186187188189190191192193194195196197198199200201202203204205206207208209210211212213214215216217218219220221222223224225226227228229230231232233234235236237238239240241242243244245246247248249250251252253254255256257258259260261262263264265266267268269270271272273274275276277278279280281282283284285286287288289290291292293294295296297298299300301302303304305306307308309310311312313314315316317318319320321322323324325326327328329330331332333334335336337338339340341342343344345346347348349350351352353354355356357358359360361362363364365366367368369370371372373374375
  1. # Copyright 2025 The HuggingFace Team. All rights reserved.
  2. #
  3. # Licensed under the Apache License, Version 2.0 (the "License");
  4. # you may not use this file except in compliance with the License.
  5. # You may obtain a copy of the License at
  6. #
  7. # http://www.apache.org/licenses/LICENSE-2.0
  8. #
  9. # Unless required by applicable law or agreed to in writing, software
  10. # distributed under the License is distributed on an "AS IS" BASIS,
  11. # WITHOUT WARRANTIES OR CONDITIONS OF ANY KIND, either express or implied.
  12. # See the License for the specific language governing permissions and
  13. # limitations under the License.
  14. import torch
  15. import torch.nn as nn
  16. import torch.nn.functional as F
  17. from ..utils import is_vision_available
  18. from .loss_for_object_detection import box_iou
  19. from .loss_rt_detr import RTDetrHungarianMatcher, RTDetrLoss
  20. if is_vision_available():
  21. from transformers.image_transforms import center_to_corners_format
  22. def _set_aux_loss(outputs_class, outputs_coord):
  23. return [{"logits": a, "pred_boxes": b} for a, b in zip(outputs_class, outputs_coord)]
  24. def _set_aux_loss2(
  25. outputs_class, outputs_coord, outputs_corners, outputs_ref, teacher_corners=None, teacher_logits=None
  26. ):
  27. return [
  28. {
  29. "logits": a,
  30. "pred_boxes": b,
  31. "pred_corners": c,
  32. "ref_points": d,
  33. "teacher_corners": teacher_corners,
  34. "teacher_logits": teacher_logits,
  35. }
  36. for a, b, c, d in zip(outputs_class, outputs_coord, outputs_corners, outputs_ref)
  37. ]
  38. def weighting_function(max_num_bins: int, up: torch.Tensor, reg_scale: int) -> torch.Tensor:
  39. """
  40. Generates the non-uniform Weighting Function W(n) for bounding box regression.
  41. Args:
  42. max_num_bins (int): Max number of the discrete bins.
  43. up (Tensor): Controls upper bounds of the sequence,
  44. where maximum offset is ±up * H / W.
  45. reg_scale (float): Controls the curvature of the Weighting Function.
  46. Larger values result in flatter weights near the central axis W(max_num_bins/2)=0
  47. and steeper weights at both ends.
  48. Returns:
  49. Tensor: Sequence of Weighting Function.
  50. """
  51. upper_bound1 = abs(up[0]) * abs(reg_scale)
  52. upper_bound2 = abs(up[0]) * abs(reg_scale) * 2
  53. step = (upper_bound1 + 1) ** (2 / (max_num_bins - 2))
  54. left_values = [-((step) ** i) + 1 for i in range(max_num_bins // 2 - 1, 0, -1)]
  55. right_values = [(step) ** i - 1 for i in range(1, max_num_bins // 2)]
  56. values = [-upper_bound2] + left_values + [torch.zeros_like(up[0][None])] + right_values + [upper_bound2]
  57. values = [v if v.dim() > 0 else v.unsqueeze(0) for v in values]
  58. values = torch.cat(values, 0)
  59. return values
  60. def translate_gt(gt: torch.Tensor, max_num_bins: int, reg_scale: int, up: torch.Tensor):
  61. """
  62. Decodes bounding box ground truth (GT) values into distribution-based GT representations.
  63. This function maps continuous GT values into discrete distribution bins, which can be used
  64. for regression tasks in object detection models. It calculates the indices of the closest
  65. bins to each GT value and assigns interpolation weights to these bins based on their proximity
  66. to the GT value.
  67. Args:
  68. gt (Tensor): Ground truth bounding box values, shape (N, ).
  69. max_num_bins (int): Maximum number of discrete bins for the distribution.
  70. reg_scale (float): Controls the curvature of the Weighting Function.
  71. up (Tensor): Controls the upper bounds of the Weighting Function.
  72. Returns:
  73. tuple[Tensor, Tensor, Tensor]:
  74. - indices (Tensor): Index of the left bin closest to each GT value, shape (N, ).
  75. - weight_right (Tensor): Weight assigned to the right bin, shape (N, ).
  76. - weight_left (Tensor): Weight assigned to the left bin, shape (N, ).
  77. """
  78. gt = gt.reshape(-1)
  79. function_values = weighting_function(max_num_bins, up, reg_scale)
  80. # Find the closest left-side indices for each value
  81. diffs = function_values.unsqueeze(0) - gt.unsqueeze(1)
  82. mask = diffs <= 0
  83. closest_left_indices = torch.sum(mask, dim=1) - 1
  84. # Calculate the weights for the interpolation
  85. indices = closest_left_indices.float()
  86. weight_right = torch.zeros_like(indices)
  87. weight_left = torch.zeros_like(indices)
  88. valid_idx_mask = (indices >= 0) & (indices < max_num_bins)
  89. valid_indices = indices[valid_idx_mask].long()
  90. # Obtain distances
  91. left_values = function_values[valid_indices]
  92. right_values = function_values[valid_indices + 1]
  93. left_diffs = torch.abs(gt[valid_idx_mask] - left_values)
  94. right_diffs = torch.abs(right_values - gt[valid_idx_mask])
  95. # Valid weights
  96. weight_right[valid_idx_mask] = left_diffs / (left_diffs + right_diffs)
  97. weight_left[valid_idx_mask] = 1.0 - weight_right[valid_idx_mask]
  98. # Invalid weights (out of range)
  99. invalid_idx_mask_neg = indices < 0
  100. weight_right[invalid_idx_mask_neg] = 0.0
  101. weight_left[invalid_idx_mask_neg] = 1.0
  102. indices[invalid_idx_mask_neg] = 0.0
  103. invalid_idx_mask_pos = indices >= max_num_bins
  104. weight_right[invalid_idx_mask_pos] = 1.0
  105. weight_left[invalid_idx_mask_pos] = 0.0
  106. indices[invalid_idx_mask_pos] = max_num_bins - 0.1
  107. return indices, weight_right, weight_left
  108. def bbox2distance(points, bbox, max_num_bins, reg_scale, up, eps=0.1):
  109. """
  110. Converts bounding box coordinates to distances from a reference point.
  111. Args:
  112. points (Tensor): (n, 4) [x, y, w, h], where (x, y) is the center.
  113. bbox (Tensor): (n, 4) bounding boxes in "xyxy" format.
  114. max_num_bins (float): Maximum bin value.
  115. reg_scale (float): Controlling curvarture of W(n).
  116. up (Tensor): Controlling upper bounds of W(n).
  117. eps (float): Small value to ensure target < max_num_bins.
  118. Returns:
  119. Tensor: Decoded distances.
  120. """
  121. reg_scale = abs(reg_scale)
  122. left = (points[:, 0] - bbox[:, 0]) / (points[..., 2] / reg_scale + 1e-16) - 0.5 * reg_scale
  123. top = (points[:, 1] - bbox[:, 1]) / (points[..., 3] / reg_scale + 1e-16) - 0.5 * reg_scale
  124. right = (bbox[:, 2] - points[:, 0]) / (points[..., 2] / reg_scale + 1e-16) - 0.5 * reg_scale
  125. bottom = (bbox[:, 3] - points[:, 1]) / (points[..., 3] / reg_scale + 1e-16) - 0.5 * reg_scale
  126. four_lens = torch.stack([left, top, right, bottom], -1)
  127. four_lens, weight_right, weight_left = translate_gt(four_lens, max_num_bins, reg_scale, up)
  128. if max_num_bins is not None:
  129. four_lens = four_lens.clamp(min=0, max=max_num_bins - eps)
  130. return four_lens.reshape(-1).detach(), weight_right.detach(), weight_left.detach()
  131. class DFineLoss(RTDetrLoss):
  132. """
  133. This class computes the losses for D-FINE. The process happens in two steps: 1) we compute hungarian assignment
  134. between ground truth boxes and the outputs of the model 2) we supervise each pair of matched ground-truth /
  135. prediction (supervise class and box).
  136. Args:
  137. matcher (`DetrHungarianMatcher`):
  138. Module able to compute a matching between targets and proposals.
  139. weight_dict (`Dict`):
  140. Dictionary relating each loss with its weights. These losses are configured in DFineConf as
  141. `weight_loss_vfl`, `weight_loss_bbox`, `weight_loss_giou`, `weight_loss_fgl`, `weight_loss_ddf`
  142. losses (`list[str]`):
  143. List of all the losses to be applied. See `get_loss` for a list of all available losses.
  144. alpha (`float`):
  145. Parameter alpha used to compute the focal loss.
  146. gamma (`float`):
  147. Parameter gamma used to compute the focal loss.
  148. eos_coef (`float`):
  149. Relative classification weight applied to the no-object category.
  150. num_classes (`int`):
  151. Number of object categories, omitting the special no-object category.
  152. """
  153. def __init__(self, config):
  154. super().__init__(config)
  155. self.matcher = RTDetrHungarianMatcher(config)
  156. self.max_num_bins = config.max_num_bins
  157. self.weight_dict = {
  158. "loss_vfl": config.weight_loss_vfl,
  159. "loss_bbox": config.weight_loss_bbox,
  160. "loss_giou": config.weight_loss_giou,
  161. "loss_fgl": config.weight_loss_fgl,
  162. "loss_ddf": config.weight_loss_ddf,
  163. }
  164. self.losses = ["vfl", "boxes", "local"]
  165. self.reg_scale = config.reg_scale
  166. self.up = nn.Parameter(torch.tensor([config.up]), requires_grad=False)
  167. def unimodal_distribution_focal_loss(
  168. self, pred, label, weight_right, weight_left, weight=None, reduction="sum", avg_factor=None
  169. ):
  170. dis_left = label.long()
  171. dis_right = dis_left + 1
  172. loss = F.cross_entropy(pred, dis_left, reduction="none") * weight_left.reshape(-1) + F.cross_entropy(
  173. pred, dis_right, reduction="none"
  174. ) * weight_right.reshape(-1)
  175. if weight is not None:
  176. weight = weight.float()
  177. loss = loss * weight
  178. if avg_factor is not None:
  179. loss = loss.sum() / avg_factor
  180. elif reduction == "mean":
  181. loss = loss.mean()
  182. elif reduction == "sum":
  183. loss = loss.sum()
  184. return loss
  185. def loss_local(self, outputs, targets, indices, num_boxes, T=5):
  186. """Compute Fine-Grained Localization (FGL) Loss
  187. and Decoupled Distillation Focal (DDF) Loss."""
  188. losses = {}
  189. if "pred_corners" in outputs:
  190. idx = self._get_source_permutation_idx(indices)
  191. target_boxes = torch.cat([t["boxes"][i] for t, (_, i) in zip(targets, indices)], dim=0)
  192. pred_corners = outputs["pred_corners"][idx].reshape(-1, (self.max_num_bins + 1))
  193. ref_points = outputs["ref_points"][idx].detach()
  194. with torch.no_grad():
  195. self.fgl_targets = bbox2distance(
  196. ref_points,
  197. center_to_corners_format(target_boxes),
  198. self.max_num_bins,
  199. self.reg_scale,
  200. self.up,
  201. )
  202. target_corners, weight_right, weight_left = self.fgl_targets
  203. ious = torch.diag(
  204. box_iou(center_to_corners_format(outputs["pred_boxes"][idx]), center_to_corners_format(target_boxes))[
  205. 0
  206. ]
  207. )
  208. weight_targets = ious.unsqueeze(-1).repeat(1, 1, 4).reshape(-1).detach()
  209. losses["loss_fgl"] = self.unimodal_distribution_focal_loss(
  210. pred_corners,
  211. target_corners,
  212. weight_right,
  213. weight_left,
  214. weight_targets,
  215. avg_factor=num_boxes,
  216. )
  217. pred_corners = outputs["pred_corners"].reshape(-1, (self.max_num_bins + 1))
  218. target_corners = outputs["teacher_corners"].reshape(-1, (self.max_num_bins + 1))
  219. if torch.equal(pred_corners, target_corners):
  220. losses["loss_ddf"] = pred_corners.sum() * 0
  221. else:
  222. weight_targets_local = outputs["teacher_logits"].sigmoid().max(dim=-1)[0]
  223. mask = torch.zeros_like(weight_targets_local, dtype=torch.bool)
  224. mask[idx] = True
  225. mask = mask.unsqueeze(-1).repeat(1, 1, 4).reshape(-1)
  226. weight_targets_local[idx] = ious.reshape_as(weight_targets_local[idx]).to(weight_targets_local.dtype)
  227. weight_targets_local = weight_targets_local.unsqueeze(-1).repeat(1, 1, 4).reshape(-1).detach()
  228. loss_match_local = (
  229. weight_targets_local
  230. * (T**2)
  231. * (
  232. nn.KLDivLoss(reduction="none")(
  233. F.log_softmax(pred_corners / T, dim=1),
  234. F.softmax(target_corners.detach() / T, dim=1),
  235. )
  236. ).sum(-1)
  237. )
  238. batch_scale = 1 / outputs["pred_boxes"].shape[0] # it should be refined
  239. self.num_pos, self.num_neg = (
  240. (mask.sum() * batch_scale) ** 0.5,
  241. ((~mask).sum() * batch_scale) ** 0.5,
  242. )
  243. loss_match_local1 = loss_match_local[mask].mean() if mask.any() else 0
  244. loss_match_local2 = loss_match_local[~mask].mean() if (~mask).any() else 0
  245. losses["loss_ddf"] = (loss_match_local1 * self.num_pos + loss_match_local2 * self.num_neg) / (
  246. self.num_pos + self.num_neg
  247. )
  248. return losses
  249. def get_loss(self, loss, outputs, targets, indices, num_boxes):
  250. loss_map = {
  251. "cardinality": self.loss_cardinality,
  252. "local": self.loss_local,
  253. "boxes": self.loss_boxes,
  254. "focal": self.loss_labels_focal,
  255. "vfl": self.loss_labels_vfl,
  256. }
  257. if loss not in loss_map:
  258. raise ValueError(f"Loss {loss} not supported")
  259. return loss_map[loss](outputs, targets, indices, num_boxes)
  260. def DFineForObjectDetectionLoss(
  261. logits,
  262. labels,
  263. device,
  264. pred_boxes,
  265. config,
  266. outputs_class=None,
  267. outputs_coord=None,
  268. enc_topk_logits=None,
  269. enc_topk_bboxes=None,
  270. denoising_meta_values=None,
  271. predicted_corners=None,
  272. initial_reference_points=None,
  273. **kwargs,
  274. ):
  275. criterion = DFineLoss(config)
  276. criterion.to(device)
  277. # Second: compute the losses, based on outputs and labels
  278. outputs_loss = {}
  279. outputs_loss["logits"] = logits
  280. outputs_loss["pred_boxes"] = pred_boxes.clamp(min=0, max=1)
  281. auxiliary_outputs = None
  282. if config.auxiliary_loss:
  283. if denoising_meta_values is not None:
  284. dn_out_coord, outputs_coord = torch.split(
  285. outputs_coord.clamp(min=0, max=1), denoising_meta_values["dn_num_split"], dim=2
  286. )
  287. dn_out_class, outputs_class = torch.split(outputs_class, denoising_meta_values["dn_num_split"], dim=2)
  288. dn_out_corners, out_corners = torch.split(predicted_corners, denoising_meta_values["dn_num_split"], dim=2)
  289. dn_out_refs, out_refs = torch.split(initial_reference_points, denoising_meta_values["dn_num_split"], dim=2)
  290. auxiliary_outputs = _set_aux_loss2(
  291. outputs_class[:, :-1].transpose(0, 1),
  292. outputs_coord[:, :-1].transpose(0, 1),
  293. out_corners[:, :-1].transpose(0, 1),
  294. out_refs[:, :-1].transpose(0, 1),
  295. out_corners[:, -1],
  296. outputs_class[:, -1],
  297. )
  298. outputs_loss["auxiliary_outputs"] = auxiliary_outputs
  299. outputs_loss["auxiliary_outputs"].extend(
  300. _set_aux_loss([enc_topk_logits], [enc_topk_bboxes.clamp(min=0, max=1)])
  301. )
  302. dn_auxiliary_outputs = _set_aux_loss2(
  303. dn_out_class.transpose(0, 1),
  304. dn_out_coord.transpose(0, 1),
  305. dn_out_corners.transpose(0, 1),
  306. dn_out_refs.transpose(0, 1),
  307. dn_out_corners[:, -1],
  308. dn_out_class[:, -1],
  309. )
  310. outputs_loss["dn_auxiliary_outputs"] = dn_auxiliary_outputs
  311. outputs_loss["denoising_meta_values"] = denoising_meta_values
  312. loss_dict = criterion(outputs_loss, labels)
  313. loss = sum(loss_dict.values())
  314. return loss, loss_dict, auxiliary_outputs