transform.py 12 KB

123456789101112131415161718192021222324252627282930313233343536373839404142434445464748495051525354555657585960616263646566676869707172737475767778798081828384858687888990919293949596979899100101102103104105106107108109110111112113114115116117118119120121122123124125126127128129130131132133134135136137138139140141142143144145146147148149150151152153154155156157158159160161162163164165166167168169170171172173174175176177178179180181182183184185186187188189190191192193194195196197198199200201202203204205206207208209210211212213214215216217218219220221222223224225226227228229230231232233234235236237238239240241242243244245246247248249250251252253254255256257258259260261262263264265266267268269270271272273274275276277278279280281282283284285286287288289290291292293294295296297298299300301302303304305306307308309310311312313314315316317318319320321322
  1. import math
  2. from typing import Any, Dict, List, Optional, Tuple
  3. import torch
  4. from torch import nn, Tensor
  5. from .image_list import ImageList
  6. from .roi_heads import paste_masks_in_image
  7. from ... import _is_tracing
  8. @torch.jit.unused
  9. def _get_shape_onnx(image: Tensor) -> Tensor:
  10. from torch.onnx import operators
  11. return operators.shape_as_tensor(image)[-2:]
  12. @torch.jit.unused
  13. def _fake_cast_onnx(v: Tensor) -> float:
  14. # ONNX requires a tensor but here we fake its type for JIT.
  15. return v
  16. def _resize_image_and_masks(
  17. image: Tensor,
  18. self_min_size: int,
  19. self_max_size: int,
  20. target: Optional[Dict[str, Tensor]] = None,
  21. fixed_size: Optional[Tuple[int, int]] = None,
  22. ) -> Tuple[Tensor, Optional[Dict[str, Tensor]]]:
  23. if _is_tracing():
  24. im_shape = _get_shape_onnx(image)
  25. elif torch.jit.is_scripting():
  26. im_shape = torch.tensor(image.shape[-2:])
  27. else:
  28. im_shape = image.shape[-2:]
  29. size: Optional[List[int]] = None
  30. scale_factor: Optional[float] = None
  31. recompute_scale_factor: Optional[bool] = None
  32. if fixed_size is not None:
  33. size = [fixed_size[1], fixed_size[0]]
  34. else:
  35. if torch.jit.is_scripting() or _is_tracing():
  36. min_size = torch.min(im_shape).to(dtype=torch.float32)
  37. max_size = torch.max(im_shape).to(dtype=torch.float32)
  38. self_min_size_f = float(self_min_size)
  39. self_max_size_f = float(self_max_size)
  40. scale = torch.min(self_min_size_f / min_size, self_max_size_f / max_size)
  41. if _is_tracing():
  42. scale_factor = _fake_cast_onnx(scale)
  43. else:
  44. scale_factor = scale.item()
  45. else:
  46. # Do it the normal way
  47. min_size = min(im_shape)
  48. max_size = max(im_shape)
  49. scale_factor = min(self_min_size / min_size, self_max_size / max_size)
  50. recompute_scale_factor = True
  51. image = torch.nn.functional.interpolate(
  52. image[None],
  53. size=size,
  54. scale_factor=scale_factor,
  55. mode="bilinear",
  56. recompute_scale_factor=recompute_scale_factor,
  57. align_corners=False,
  58. )[0]
  59. if target is None:
  60. return image, target
  61. if "masks" in target:
  62. mask = target["masks"]
  63. mask = torch.nn.functional.interpolate(
  64. mask[:, None].float(), size=size, scale_factor=scale_factor, recompute_scale_factor=recompute_scale_factor
  65. )[:, 0].byte()
  66. target["masks"] = mask
  67. return image, target
  68. class GeneralizedRCNNTransform(nn.Module):
  69. """
  70. Performs input / target transformation before feeding the data to a GeneralizedRCNN
  71. model.
  72. The transformations it performs are:
  73. - input normalization (mean subtraction and std division)
  74. - input / target resizing to match min_size / max_size
  75. It returns a ImageList for the inputs, and a List[Dict[Tensor]] for the targets
  76. """
  77. def __init__(
  78. self,
  79. min_size: int,
  80. max_size: int,
  81. image_mean: List[float],
  82. image_std: List[float],
  83. size_divisible: int = 32,
  84. fixed_size: Optional[Tuple[int, int]] = None,
  85. **kwargs: Any,
  86. ):
  87. super().__init__()
  88. if not isinstance(min_size, (list, tuple)):
  89. min_size = (min_size,)
  90. self.min_size = min_size
  91. self.max_size = max_size
  92. self.image_mean = image_mean
  93. self.image_std = image_std
  94. self.size_divisible = size_divisible
  95. self.fixed_size = fixed_size
  96. self._skip_resize = kwargs.pop("_skip_resize", False)
  97. def forward(
  98. self, images: List[Tensor], targets: Optional[List[Dict[str, Tensor]]] = None
  99. ) -> Tuple[ImageList, Optional[List[Dict[str, Tensor]]]]:
  100. images = [img for img in images]
  101. if targets is not None:
  102. # make a copy of targets to avoid modifying it in-place
  103. # once torchscript supports dict comprehension
  104. # this can be simplified as follows
  105. # targets = [{k: v for k,v in t.items()} for t in targets]
  106. targets_copy: List[Dict[str, Tensor]] = []
  107. for t in targets:
  108. data: Dict[str, Tensor] = {}
  109. for k, v in t.items():
  110. data[k] = v
  111. targets_copy.append(data)
  112. targets = targets_copy
  113. for i in range(len(images)):
  114. image = images[i]
  115. target_index = targets[i] if targets is not None else None
  116. if image.dim() != 3:
  117. raise ValueError(f"images is expected to be a list of 3d tensors of shape [C, H, W], got {image.shape}")
  118. image = self.normalize(image)
  119. image, target_index = self.resize(image, target_index)
  120. images[i] = image
  121. if targets is not None and target_index is not None:
  122. targets[i] = target_index
  123. image_sizes = [img.shape[-2:] for img in images]
  124. images = self.batch_images(images, size_divisible=self.size_divisible)
  125. image_sizes_list: List[Tuple[int, int]] = []
  126. for image_size in image_sizes:
  127. torch._assert(
  128. len(image_size) == 2,
  129. f"Input tensors expected to have in the last two elements H and W, instead got {image_size}",
  130. )
  131. image_sizes_list.append((image_size[0], image_size[1]))
  132. image_list = ImageList(images, image_sizes_list)
  133. return image_list, targets
  134. def normalize(self, image: Tensor) -> Tensor:
  135. if not image.is_floating_point():
  136. raise TypeError(
  137. f"Expected input images to be of floating type (in range [0, 1]), "
  138. f"but found type {image.dtype} instead"
  139. )
  140. dtype, device = image.dtype, image.device
  141. mean = torch.as_tensor(self.image_mean, dtype=dtype, device=device)
  142. std = torch.as_tensor(self.image_std, dtype=dtype, device=device)
  143. # print(f'mean:{mean}')
  144. # print(f'std:{std}')
  145. # print(f'image:{image.shape}')
  146. return (image - mean[:, None, None]) / std[:, None, None]
  147. def torch_choice(self, k: List[int]) -> int:
  148. """
  149. Implements `random.choice` via torch ops, so it can be compiled with
  150. TorchScript and we use PyTorch's RNG (not native RNG)
  151. """
  152. index = int(torch.empty(1).uniform_(0.0, float(len(k))).item())
  153. return k[index]
  154. def resize(
  155. self,
  156. image: Tensor,
  157. target: Optional[Dict[str, Tensor]] = None,
  158. ) -> Tuple[Tensor, Optional[Dict[str, Tensor]]]:
  159. h, w = image.shape[-2:]
  160. if self.training:
  161. if self._skip_resize:
  162. return image, target
  163. size = self.torch_choice(self.min_size)
  164. else:
  165. size = self.min_size[-1]
  166. image, target = _resize_image_and_masks(image, size, self.max_size, target, self.fixed_size)
  167. if target is None:
  168. return image, target
  169. bbox = target["boxes"]
  170. bbox = resize_boxes(bbox, (h, w), image.shape[-2:])
  171. target["boxes"] = bbox
  172. if "keypoints" in target:
  173. keypoints = target["keypoints"]
  174. keypoints = resize_keypoints(keypoints, (h, w), image.shape[-2:])
  175. target["keypoints"] = keypoints
  176. return image, target
  177. # _onnx_batch_images() is an implementation of
  178. # batch_images() that is supported by ONNX tracing.
  179. @torch.jit.unused
  180. def _onnx_batch_images(self, images: List[Tensor], size_divisible: int = 32) -> Tensor:
  181. max_size = []
  182. for i in range(images[0].dim()):
  183. max_size_i = torch.max(torch.stack([img.shape[i] for img in images]).to(torch.float32)).to(torch.int64)
  184. max_size.append(max_size_i)
  185. stride = size_divisible
  186. max_size[1] = (torch.ceil((max_size[1].to(torch.float32)) / stride) * stride).to(torch.int64)
  187. max_size[2] = (torch.ceil((max_size[2].to(torch.float32)) / stride) * stride).to(torch.int64)
  188. max_size = tuple(max_size)
  189. # work around for
  190. # pad_img[: img.shape[0], : img.shape[1], : img.shape[2]].copy_(img)
  191. # which is not yet supported in onnx
  192. padded_imgs = []
  193. for img in images:
  194. padding = [(s1 - s2) for s1, s2 in zip(max_size, tuple(img.shape))]
  195. padded_img = torch.nn.functional.pad(img, (0, padding[2], 0, padding[1], 0, padding[0]))
  196. padded_imgs.append(padded_img)
  197. return torch.stack(padded_imgs)
  198. def max_by_axis(self, the_list: List[List[int]]) -> List[int]:
  199. maxes = the_list[0]
  200. for sublist in the_list[1:]:
  201. for index, item in enumerate(sublist):
  202. maxes[index] = max(maxes[index], item)
  203. return maxes
  204. def batch_images(self, images: List[Tensor], size_divisible: int = 32) -> Tensor:
  205. if _is_tracing():
  206. # batch_images() does not export well to ONNX
  207. # call _onnx_batch_images() instead
  208. return self._onnx_batch_images(images, size_divisible)
  209. max_size = self.max_by_axis([list(img.shape) for img in images])
  210. stride = float(size_divisible)
  211. max_size = list(max_size)
  212. max_size[1] = int(math.ceil(float(max_size[1]) / stride) * stride)
  213. max_size[2] = int(math.ceil(float(max_size[2]) / stride) * stride)
  214. batch_shape = [len(images)] + max_size
  215. batched_imgs = images[0].new_full(batch_shape, 0)
  216. for i in range(batched_imgs.shape[0]):
  217. img = images[i]
  218. batched_imgs[i, : img.shape[0], : img.shape[1], : img.shape[2]].copy_(img)
  219. return batched_imgs
  220. def postprocess(
  221. self,
  222. result: List[Dict[str, Tensor]],
  223. image_shapes: List[Tuple[int, int]],
  224. original_image_sizes: List[Tuple[int, int]],
  225. ) -> List[Dict[str, Tensor]]:
  226. if self.training:
  227. return result
  228. for i, (pred, im_s, o_im_s) in enumerate(zip(result, image_shapes, original_image_sizes)):
  229. boxes = pred["boxes"]
  230. boxes = resize_boxes(boxes, im_s, o_im_s)
  231. result[i]["boxes"] = boxes
  232. if "masks" in pred:
  233. masks = pred["masks"]
  234. masks = paste_masks_in_image(masks, boxes, o_im_s)
  235. result[i]["masks"] = masks
  236. if "keypoints" in pred:
  237. keypoints = pred["keypoints"]
  238. keypoints = resize_keypoints(keypoints, im_s, o_im_s)
  239. result[i]["keypoints"] = keypoints
  240. return result
  241. def __repr__(self) -> str:
  242. format_string = f"{self.__class__.__name__}("
  243. _indent = "\n "
  244. format_string += f"{_indent}Normalize(mean={self.image_mean}, std={self.image_std})"
  245. format_string += f"{_indent}Resize(min_size={self.min_size}, max_size={self.max_size}, mode='bilinear')"
  246. format_string += "\n)"
  247. return format_string
  248. def resize_keypoints(keypoints: Tensor, original_size: List[int], new_size: List[int]) -> Tensor:
  249. ratios = [
  250. torch.tensor(s, dtype=torch.float32, device=keypoints.device)
  251. / torch.tensor(s_orig, dtype=torch.float32, device=keypoints.device)
  252. for s, s_orig in zip(new_size, original_size)
  253. ]
  254. ratio_h, ratio_w = ratios
  255. resized_data = keypoints.clone()
  256. if torch._C._get_tracing_state():
  257. resized_data_0 = resized_data[:, :, 0] * ratio_w
  258. resized_data_1 = resized_data[:, :, 1] * ratio_h
  259. resized_data = torch.stack((resized_data_0, resized_data_1, resized_data[:, :, 2]), dim=2)
  260. else:
  261. resized_data[..., 0] *= ratio_w
  262. resized_data[..., 1] *= ratio_h
  263. return resized_data
  264. def resize_boxes(boxes: Tensor, original_size: List[int], new_size: List[int]) -> Tensor:
  265. ratios = [
  266. torch.tensor(s, dtype=torch.float32, device=boxes.device)
  267. / torch.tensor(s_orig, dtype=torch.float32, device=boxes.device)
  268. for s, s_orig in zip(new_size, original_size)
  269. ]
  270. ratio_height, ratio_width = ratios
  271. xmin, ymin, xmax, ymax = boxes.unbind(1)
  272. xmin = xmin * ratio_width
  273. xmax = xmax * ratio_width
  274. ymin = ymin * ratio_height
  275. ymax = ymax * ratio_height
  276. return torch.stack((xmin, ymin, xmax, ymax), dim=1)