simota_assigner.py 11 KB

123456789101112131415161718192021222324252627282930313233343536373839404142434445464748495051525354555657585960616263646566676869707172737475767778798081828384858687888990919293949596979899100101102103104105106107108109110111112113114115116117118119120121122123124125126127128129130131132133134135136137138139140141142143144145146147148149150151152153154155156157158159160161162163164165166167168169170171172173174175176177178179180181182183184185186187188189190191192193194195196197198199200201202203204205206207208209210211212213214215216217218219220221222223224225226227228229230231232233234235236237238239240241242243244245246247248249250251252253254255256257258259260261262
  1. # Copyright (c) 2021 PaddlePaddle Authors. 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. # The code is based on:
  15. # https://github.com/open-mmlab/mmdetection/blob/master/mmdet/core/bbox/assigners/sim_ota_assigner.py
  16. import paddle
  17. import numpy as np
  18. import paddle.nn.functional as F
  19. from ppdet.modeling.losses.varifocal_loss import varifocal_loss
  20. from ppdet.modeling.bbox_utils import batch_bbox_overlaps
  21. from ppdet.core.workspace import register
  22. @register
  23. class SimOTAAssigner(object):
  24. """Computes matching between predictions and ground truth.
  25. Args:
  26. center_radius (int | float, optional): Ground truth center size
  27. to judge whether a prior is in center. Default 2.5.
  28. candidate_topk (int, optional): The candidate top-k which used to
  29. get top-k ious to calculate dynamic-k. Default 10.
  30. iou_weight (int | float, optional): The scale factor for regression
  31. iou cost. Default 3.0.
  32. cls_weight (int | float, optional): The scale factor for classification
  33. cost. Default 1.0.
  34. num_classes (int): The num_classes of dataset.
  35. use_vfl (int): Whether to use varifocal_loss when calculating the cost matrix.
  36. """
  37. __shared__ = ['num_classes']
  38. def __init__(self,
  39. center_radius=2.5,
  40. candidate_topk=10,
  41. iou_weight=3.0,
  42. cls_weight=1.0,
  43. num_classes=80,
  44. use_vfl=True):
  45. self.center_radius = center_radius
  46. self.candidate_topk = candidate_topk
  47. self.iou_weight = iou_weight
  48. self.cls_weight = cls_weight
  49. self.num_classes = num_classes
  50. self.use_vfl = use_vfl
  51. def get_in_gt_and_in_center_info(self, flatten_center_and_stride,
  52. gt_bboxes):
  53. num_gt = gt_bboxes.shape[0]
  54. flatten_x = flatten_center_and_stride[:, 0].unsqueeze(1).tile(
  55. [1, num_gt])
  56. flatten_y = flatten_center_and_stride[:, 1].unsqueeze(1).tile(
  57. [1, num_gt])
  58. flatten_stride_x = flatten_center_and_stride[:, 2].unsqueeze(1).tile(
  59. [1, num_gt])
  60. flatten_stride_y = flatten_center_and_stride[:, 3].unsqueeze(1).tile(
  61. [1, num_gt])
  62. # is prior centers in gt bboxes, shape: [n_center, n_gt]
  63. l_ = flatten_x - gt_bboxes[:, 0]
  64. t_ = flatten_y - gt_bboxes[:, 1]
  65. r_ = gt_bboxes[:, 2] - flatten_x
  66. b_ = gt_bboxes[:, 3] - flatten_y
  67. deltas = paddle.stack([l_, t_, r_, b_], axis=1)
  68. is_in_gts = deltas.min(axis=1) > 0
  69. is_in_gts_all = is_in_gts.sum(axis=1) > 0
  70. # is prior centers in gt centers
  71. gt_center_xs = (gt_bboxes[:, 0] + gt_bboxes[:, 2]) / 2.0
  72. gt_center_ys = (gt_bboxes[:, 1] + gt_bboxes[:, 3]) / 2.0
  73. ct_bound_l = gt_center_xs - self.center_radius * flatten_stride_x
  74. ct_bound_t = gt_center_ys - self.center_radius * flatten_stride_y
  75. ct_bound_r = gt_center_xs + self.center_radius * flatten_stride_x
  76. ct_bound_b = gt_center_ys + self.center_radius * flatten_stride_y
  77. cl_ = flatten_x - ct_bound_l
  78. ct_ = flatten_y - ct_bound_t
  79. cr_ = ct_bound_r - flatten_x
  80. cb_ = ct_bound_b - flatten_y
  81. ct_deltas = paddle.stack([cl_, ct_, cr_, cb_], axis=1)
  82. is_in_cts = ct_deltas.min(axis=1) > 0
  83. is_in_cts_all = is_in_cts.sum(axis=1) > 0
  84. # in any of gts or gt centers, shape: [n_center]
  85. is_in_gts_or_centers_all = paddle.logical_or(is_in_gts_all,
  86. is_in_cts_all)
  87. is_in_gts_or_centers_all_inds = paddle.nonzero(
  88. is_in_gts_or_centers_all).squeeze(1)
  89. # both in gts and gt centers, shape: [num_fg, num_gt]
  90. is_in_gts_and_centers = paddle.logical_and(
  91. paddle.gather(
  92. is_in_gts.cast('int'), is_in_gts_or_centers_all_inds,
  93. axis=0).cast('bool'),
  94. paddle.gather(
  95. is_in_cts.cast('int'), is_in_gts_or_centers_all_inds,
  96. axis=0).cast('bool'))
  97. return is_in_gts_or_centers_all, is_in_gts_or_centers_all_inds, is_in_gts_and_centers
  98. def dynamic_k_matching(self, cost_matrix, pairwise_ious, num_gt):
  99. match_matrix = np.zeros_like(cost_matrix.numpy())
  100. # select candidate topk ious for dynamic-k calculation
  101. topk_ious, _ = paddle.topk(pairwise_ious, self.candidate_topk, axis=0)
  102. # calculate dynamic k for each gt
  103. dynamic_ks = paddle.clip(topk_ious.sum(0).cast('int'), min=1)
  104. for gt_idx in range(num_gt):
  105. _, pos_idx = paddle.topk(
  106. cost_matrix[:, gt_idx], k=dynamic_ks[gt_idx], largest=False)
  107. match_matrix[:, gt_idx][pos_idx.numpy()] = 1.0
  108. del topk_ious, dynamic_ks, pos_idx
  109. # match points more than two gts
  110. extra_match_gts_mask = match_matrix.sum(1) > 1
  111. if extra_match_gts_mask.sum() > 0:
  112. cost_matrix = cost_matrix.numpy()
  113. cost_argmin = np.argmin(
  114. cost_matrix[extra_match_gts_mask, :], axis=1)
  115. match_matrix[extra_match_gts_mask, :] *= 0.0
  116. match_matrix[extra_match_gts_mask, cost_argmin] = 1.0
  117. # get foreground mask
  118. match_fg_mask_inmatrix = match_matrix.sum(1) > 0
  119. match_gt_inds_to_fg = match_matrix[match_fg_mask_inmatrix, :].argmax(1)
  120. return match_gt_inds_to_fg, match_fg_mask_inmatrix
  121. def get_sample(self, assign_gt_inds, gt_bboxes):
  122. pos_inds = np.unique(np.nonzero(assign_gt_inds > 0)[0])
  123. neg_inds = np.unique(np.nonzero(assign_gt_inds == 0)[0])
  124. pos_assigned_gt_inds = assign_gt_inds[pos_inds] - 1
  125. if gt_bboxes.size == 0:
  126. # hack for index error case
  127. assert pos_assigned_gt_inds.size == 0
  128. pos_gt_bboxes = np.empty_like(gt_bboxes).reshape(-1, 4)
  129. else:
  130. if len(gt_bboxes.shape) < 2:
  131. gt_bboxes = gt_bboxes.resize(-1, 4)
  132. pos_gt_bboxes = gt_bboxes[pos_assigned_gt_inds, :]
  133. return pos_inds, neg_inds, pos_gt_bboxes, pos_assigned_gt_inds
  134. def __call__(self,
  135. flatten_cls_pred_scores,
  136. flatten_center_and_stride,
  137. flatten_bboxes,
  138. gt_bboxes,
  139. gt_labels,
  140. eps=1e-7):
  141. """Assign gt to priors using SimOTA.
  142. TODO: add comment.
  143. Returns:
  144. assign_result: The assigned result.
  145. """
  146. num_gt = gt_bboxes.shape[0]
  147. num_bboxes = flatten_bboxes.shape[0]
  148. if num_gt == 0 or num_bboxes == 0:
  149. # No ground truth or boxes
  150. label = np.ones([num_bboxes], dtype=np.int64) * self.num_classes
  151. label_weight = np.ones([num_bboxes], dtype=np.float32)
  152. bbox_target = np.zeros_like(flatten_center_and_stride)
  153. return 0, label, label_weight, bbox_target
  154. is_in_gts_or_centers_all, is_in_gts_or_centers_all_inds, is_in_boxes_and_center = self.get_in_gt_and_in_center_info(
  155. flatten_center_and_stride, gt_bboxes)
  156. # bboxes and scores to calculate matrix
  157. valid_flatten_bboxes = flatten_bboxes[is_in_gts_or_centers_all_inds]
  158. valid_cls_pred_scores = flatten_cls_pred_scores[
  159. is_in_gts_or_centers_all_inds]
  160. num_valid_bboxes = valid_flatten_bboxes.shape[0]
  161. pairwise_ious = batch_bbox_overlaps(valid_flatten_bboxes,
  162. gt_bboxes) # [num_points,num_gts]
  163. if self.use_vfl:
  164. gt_vfl_labels = gt_labels.squeeze(-1).unsqueeze(0).tile(
  165. [num_valid_bboxes, 1]).reshape([-1])
  166. valid_pred_scores = valid_cls_pred_scores.unsqueeze(1).tile(
  167. [1, num_gt, 1]).reshape([-1, self.num_classes])
  168. vfl_score = np.zeros(valid_pred_scores.shape)
  169. vfl_score[np.arange(0, vfl_score.shape[0]), gt_vfl_labels.numpy(
  170. )] = pairwise_ious.reshape([-1])
  171. vfl_score = paddle.to_tensor(vfl_score)
  172. losses_vfl = varifocal_loss(
  173. valid_pred_scores, vfl_score,
  174. use_sigmoid=False).reshape([num_valid_bboxes, num_gt])
  175. losses_giou = batch_bbox_overlaps(
  176. valid_flatten_bboxes, gt_bboxes, mode='giou')
  177. cost_matrix = (
  178. losses_vfl * self.cls_weight + losses_giou * self.iou_weight +
  179. paddle.logical_not(is_in_boxes_and_center).cast('float32') *
  180. 100000000)
  181. else:
  182. iou_cost = -paddle.log(pairwise_ious + eps)
  183. gt_onehot_label = (F.one_hot(
  184. gt_labels.squeeze(-1).cast(paddle.int64),
  185. flatten_cls_pred_scores.shape[-1]).cast('float32').unsqueeze(0)
  186. .tile([num_valid_bboxes, 1, 1]))
  187. valid_pred_scores = valid_cls_pred_scores.unsqueeze(1).tile(
  188. [1, num_gt, 1])
  189. cls_cost = F.binary_cross_entropy(
  190. valid_pred_scores, gt_onehot_label, reduction='none').sum(-1)
  191. cost_matrix = (
  192. cls_cost * self.cls_weight + iou_cost * self.iou_weight +
  193. paddle.logical_not(is_in_boxes_and_center).cast('float32') *
  194. 100000000)
  195. match_gt_inds_to_fg, match_fg_mask_inmatrix = \
  196. self.dynamic_k_matching(
  197. cost_matrix, pairwise_ious, num_gt)
  198. # sample and assign results
  199. assigned_gt_inds = np.zeros([num_bboxes], dtype=np.int64)
  200. match_fg_mask_inall = np.zeros_like(assigned_gt_inds)
  201. match_fg_mask_inall[is_in_gts_or_centers_all.numpy(
  202. )] = match_fg_mask_inmatrix
  203. assigned_gt_inds[match_fg_mask_inall.astype(
  204. np.bool)] = match_gt_inds_to_fg + 1
  205. pos_inds, neg_inds, pos_gt_bboxes, pos_assigned_gt_inds \
  206. = self.get_sample(assigned_gt_inds, gt_bboxes.numpy())
  207. bbox_target = np.zeros_like(flatten_bboxes)
  208. bbox_weight = np.zeros_like(flatten_bboxes)
  209. label = np.ones([num_bboxes], dtype=np.int64) * self.num_classes
  210. label_weight = np.zeros([num_bboxes], dtype=np.float32)
  211. if len(pos_inds) > 0:
  212. gt_labels = gt_labels.numpy()
  213. pos_bbox_targets = pos_gt_bboxes
  214. bbox_target[pos_inds, :] = pos_bbox_targets
  215. bbox_weight[pos_inds, :] = 1.0
  216. if not np.any(gt_labels):
  217. label[pos_inds] = 0
  218. else:
  219. label[pos_inds] = gt_labels.squeeze(-1)[pos_assigned_gt_inds]
  220. label_weight[pos_inds] = 1.0
  221. if len(neg_inds) > 0:
  222. label_weight[neg_inds] = 1.0
  223. pos_num = max(pos_inds.size, 1)
  224. return pos_num, label, label_weight, bbox_target