_pattern_matcher.py 24 KB

123456789101112131415161718192021222324252627282930313233343536373839404142434445464748495051525354555657585960616263646566676869707172737475767778798081828384858687888990919293949596979899100101102103104105106107108109110111112113114115116117118119120121122123124125126127128129130131132133134135136137138139140141142143144145146147148149150151152153154155156157158159160161162163164165166167168169170171172173174175176177178179180181182183184185186187188189190191192193194195196197198199200201202203204205206207208209210211212213214215216217218219220221222223224225226227228229230231232233234235236237238239240241242243244245246247248249250251252253254255256257258259260261262263264265266267268269270271272273274275276277278279280281282283284285286287288289290291292293294295296297298299300301302303304305306307308309310311312313314315316317318319320321322323324325326327328329330331332333334335336337338339340341342343344345346347348349350351352353354355356357358359360361362363364365366367368369370371372373374375376377378379380381382383384385386387388389390391392393394395396397398399400401402403404405406407408409410411412413414415416417418419420421422423424425426427428429430431432433434435436437438439440441442443444445446447448449450451452453454455456457458459460461462463464465466467468469470471472473474475476477478479480481482483484485486487488489490491492493494495496497498499500501502503504505506507508509510511512513514515516517518519520521522523524525526527528529530531532533534535536537538539540541542543544545546547548549550551552553554555556557558559560561562563564565566567568569570571572573574575576577578579580581582583584585586587588589590591592593594595596597598599600601602603604605606607608609610611612613614615616617618619620621622623624625626627628629630631632633634635636637638639640641642643644645646647648649650651652653654655656657658659660661662663664665666667
  1. # mypy: allow-untyped-defs
  2. import json
  3. import math
  4. import os
  5. import re
  6. from typing import Optional
  7. import torch
  8. import torch.utils.benchmark as benchmark
  9. from torch._C._profiler import (
  10. _EventType,
  11. _ExtraFields_PyCall,
  12. _ExtraFields_PyCCall,
  13. _ExtraFields_TorchOp,
  14. _ProfilerEvent,
  15. )
  16. from torch.profiler import profile
  17. from torch.profiler._utils import index_of_first_match, traverse_bfs, traverse_dfs
  18. class Pattern:
  19. """
  20. Base class for all patterns, subclass this class and implement match()
  21. to define custom patterns.
  22. In subclass, define description and skip property.
  23. """
  24. def __init__(self, prof: profile, should_benchmark: bool = False) -> None:
  25. self.prof = prof
  26. self.should_benchmark = should_benchmark
  27. self.name = "Please specify a name for pattern"
  28. self.description = "Please specify a description for pattern"
  29. self.url = ""
  30. assert prof.profiler is not None and prof.profiler.kineto_results is not None
  31. self.event_tree = prof.profiler.kineto_results.experimental_event_tree()
  32. self.tid_root: dict[int, list[_ProfilerEvent]] = {}
  33. for event in self.event_tree:
  34. self.tid_root.setdefault(event.start_tid, []).append(event)
  35. @property
  36. def skip(self) -> bool:
  37. return False
  38. def report(self, event: _ProfilerEvent):
  39. msg = (
  40. f"{self.description}\n[Source Code Location] {source_code_location(event)}"
  41. )
  42. return msg
  43. def eventTreeTraversal(self):
  44. """
  45. Traverse the event tree and yield all events.
  46. Override this method in subclass to customize the traversal.
  47. """
  48. yield from traverse_dfs(self.event_tree)
  49. def summary(self, events: list[_ProfilerEvent]):
  50. default_summary = f"{self.name}: {len(events)} events matched."
  51. if self.should_benchmark:
  52. # If benchmark summary is not empty, use it.
  53. return (
  54. self.benchmark_summary(events)
  55. if hasattr(self, "benchmark") # type: ignore[attr-defined]
  56. else default_summary
  57. )
  58. return default_summary
  59. def benchmark_summary(self, events: list[_ProfilerEvent]) -> str:
  60. def format_time(time_ns: int) -> str:
  61. unit_lst = ["ns", "us", "ms"]
  62. for unit in unit_lst:
  63. if time_ns < 1000:
  64. return f"{time_ns:.2f} {unit}"
  65. time_ns //= 1000
  66. return f"{time_ns:.2f} s"
  67. assert hasattr(self, "benchmark"), "Please implement benchmark()"
  68. shapes_factor_map = self.benchmark(events) # type: ignore[attr-defined]
  69. original_time = sum(event.duration_time_ns for event in events)
  70. new_time = sum(
  71. shapes_factor_map[input_shapes(event)] * event.duration_time_ns
  72. for event in events
  73. )
  74. return (
  75. f"{self.name}: {len(events)} events matched. "
  76. f"Total Estimated Speedup: {format_time(original_time - new_time)} ({round(original_time / new_time, 2)}X)"
  77. )
  78. def match(self, event: _ProfilerEvent):
  79. """
  80. Return True if the event matches the pattern.
  81. This method should be overridden in subclass.
  82. """
  83. raise NotImplementedError
  84. def matched_events(self):
  85. if self.skip:
  86. return []
  87. matched_events = [
  88. event for event in self.eventTreeTraversal() if self.match(event)
  89. ]
  90. return matched_events
  91. def root_of(self, event: _ProfilerEvent):
  92. while event.parent:
  93. event = event.parent
  94. return event
  95. def siblings_of(self, event: _ProfilerEvent):
  96. if event.parent:
  97. children = event.parent.children
  98. else:
  99. children = self.tid_root[event.start_tid]
  100. index = children.index(event)
  101. return children[:index], children[index + 1 :]
  102. def next_of(self, event: _ProfilerEvent):
  103. _, next_events = self.siblings_of(event)
  104. return next_events[0] if next_events else None
  105. def prev_of(self, event: _ProfilerEvent):
  106. prev_events, _ = self.siblings_of(event)
  107. return prev_events[-1] if prev_events else None
  108. def go_up_until(self, event: _ProfilerEvent, predicate):
  109. if not event:
  110. return None
  111. while event.parent and not predicate(event):
  112. event = event.parent
  113. return event
  114. # Patterns
  115. class NamePattern(Pattern):
  116. def __init__(
  117. self, prof: profile, name: str, should_benchmark: bool = False
  118. ) -> None:
  119. super().__init__(prof, should_benchmark)
  120. self.description = f"Matched Name Event: {name}"
  121. self.name = name
  122. def match(self, event: _ProfilerEvent):
  123. return re.search(self.name, event.name) is not None
  124. class ExtraCUDACopyPattern(Pattern):
  125. """
  126. This pattern identifies if we creates a constant tensor on CPU and immediately moves it to GPU.
  127. example: torch.zeros((100, 100)).to("cuda")
  128. Pattern:
  129. built-in method |built-in method
  130. ... | aten::to
  131. aten::fill_/aten::zero_ | aten::_to_copy
  132. Algorithm:
  133. We start at node aten::to, go parent events' previous events,
  134. and check if we have a aten::fill_/aten::zero_ as we keep going down the tree.
  135. We always select the last child in the children list when we go down the tree.
  136. If at any step we failed, it is not a match.
  137. """
  138. def __init__(self, prof: profile, should_benchmark: bool = False) -> None:
  139. super().__init__(prof, should_benchmark)
  140. self.name = "Extra CUDA Copy Pattern"
  141. self.description = "Filled a CPU tensor and immediately moved it to GPU. Please initialize it on GPU."
  142. self.url = "https://pytorch.org/tutorials/recipes/recipes/tuning_guide.html#create-tensors-directly-on-the-target-device"
  143. self.init_ops = {
  144. "aten::fill_",
  145. "aten::zero_",
  146. "aten::normal_",
  147. "aten::uniform_",
  148. }
  149. @property
  150. def skip(self) -> bool:
  151. return not self.prof.with_stack or not self.prof.record_shapes
  152. def match(self, event):
  153. # TODO: We should also check tensor identities
  154. if event.name != "aten::to":
  155. return False
  156. to_event = event
  157. if not event.children:
  158. return False
  159. event = event.children[-1]
  160. if event.name != "aten::_to_copy":
  161. return False
  162. if not event.children:
  163. return False
  164. event = event.children[-1]
  165. if event.name != "aten::copy_":
  166. return False
  167. # aten::copy_ should have the first 2 args dtype the same
  168. dtypes = input_dtypes(event)
  169. if len(dtypes) < 2:
  170. return False
  171. if dtypes[0] is None or dtypes[0] != dtypes[1]:
  172. return False
  173. event = to_event
  174. # Up one level
  175. event = event.parent
  176. if event is None:
  177. return False
  178. # Check if we have a aten::fill_ in previous leaf
  179. event = self.prev_of(event)
  180. if event is None:
  181. return False
  182. while event.children:
  183. event = event.children[-1]
  184. # aten::zero_ is a special optimization case where fill_ is not called
  185. if event.name in self.init_ops:
  186. return True
  187. return event.name in self.init_ops
  188. # TODO: Check if tensor is reused
  189. def benchmark(self, events: list[_ProfilerEvent]):
  190. shapes_factor_map = {input_shapes(event): 0.0 for event in events}
  191. for shape in shapes_factor_map:
  192. size = shape[0]
  193. to_timer = benchmark.Timer(
  194. stmt='torch.ones(size).to("cuda")', globals={"size": size}
  195. )
  196. de_timer = benchmark.Timer(
  197. stmt='torch.ones(size, device="cuda")', globals={"size": size}
  198. )
  199. to_time = to_timer.timeit(10).mean
  200. de_time = de_timer.timeit(10).mean
  201. shapes_factor_map[shape] = de_time / to_time
  202. return shapes_factor_map
  203. class ForLoopIndexingPattern(Pattern):
  204. """
  205. This pattern identifies if we use a for loop to index a tensor that
  206. can be vectorized.
  207. example:
  208. tensor = torch.empty((100, 100))
  209. for i in range(100):
  210. tensor[i] = i
  211. Pattern:
  212. aten::select | ... | aten::select | ... (Repeat)
  213. Algorithm:
  214. We start at node aten::select, and we check if we can find this alternating patterns.
  215. We also keep a dictionary to avoid duplicate match in the for loop.
  216. """
  217. def __init__(self, prof: profile, should_benchmark: bool = False) -> None:
  218. super().__init__(prof, should_benchmark)
  219. self.name = "For Loop Indexing Pattern"
  220. self.description = "For loop indexing detected. Vectorization recommended."
  221. self.visited: set[int] = set()
  222. def eventTreeTraversal(self):
  223. """
  224. We need to use BFS traversal order to avoid duplicate match.
  225. """
  226. yield from traverse_bfs(self.event_tree)
  227. def match(self, event: _ProfilerEvent):
  228. if event.name != "aten::select":
  229. return False
  230. if event.id in self.visited:
  231. return False
  232. repeat_count = 1
  233. _, next = self.siblings_of(event)
  234. if len(next) <= 1:
  235. return False
  236. # Custom event list matching
  237. def same_ops(list1, list2) -> bool:
  238. if len(list1) != len(list2):
  239. return False
  240. for op1, op2 in zip(list1, list2):
  241. if op1.name != op2.name:
  242. return False
  243. return True
  244. # Record the ops between two aten::select
  245. next_select_idx = index_of_first_match(next, lambda e: e.name == "aten::select")
  246. if next_select_idx is None:
  247. return False
  248. indexing_ops = [event] + next[:next_select_idx]
  249. next = next[len(indexing_ops) - 1 :]
  250. for i in range(0, len(next), len(indexing_ops)):
  251. if same_ops(indexing_ops, next[i : i + len(indexing_ops)]):
  252. repeat_count += 1
  253. self.visited.add(next[i].id)
  254. else:
  255. break
  256. return repeat_count >= 10
  257. class FP32MatMulPattern(Pattern):
  258. def __init__(self, prof: profile, should_benchmark: bool = False) -> None:
  259. super().__init__(prof, should_benchmark)
  260. self.name = "FP32 MatMul Pattern"
  261. self.description = (
  262. "You are currently using GPU that supports TF32. "
  263. "Please enable TF32 by setting 'torch.backends.cuda.matmul.allow_tf32 = True'"
  264. )
  265. self.url = "https://pytorch.org/docs/stable/notes/cuda.html#tensorfloat-32-tf32-on-ampere-devices"
  266. @property
  267. def skip(self):
  268. if torch.version.hip is not None:
  269. has_tf32 = False
  270. else:
  271. # Anything less than sm_80 is not Ampere which doesn't support TF32
  272. has_tf32 = all(
  273. int(re.sub("sm_|compute_", "", arch)) >= 80
  274. for arch in torch.cuda.get_arch_list()
  275. )
  276. return has_tf32 is False or super().skip or not self.prof.record_shapes
  277. def match(self, event: _ProfilerEvent) -> bool:
  278. # If we saw this pattern once, we don't need to match it again
  279. if event.tag != _EventType.TorchOp:
  280. return False
  281. assert isinstance(event.extra_fields, _ExtraFields_TorchOp)
  282. if event.name == "aten::mm":
  283. if event.extra_fields.allow_tf32_cublas is False:
  284. return True
  285. return False
  286. def report(self, event: _ProfilerEvent):
  287. return self.description
  288. def benchmark(self, events: list[_ProfilerEvent]):
  289. shapes_factor_map = {input_shapes(event): 0.0 for event in events}
  290. for shape in shapes_factor_map:
  291. matrixA = torch.randn(shape[0], device="cuda", dtype=torch.float32)
  292. matrixB = torch.randn(shape[1], device="cuda", dtype=torch.float32)
  293. fp32_timer = benchmark.Timer(
  294. stmt="torch.mm(matrixA, matrixB)",
  295. globals={"matrixA": matrixA, "matrixB": matrixB},
  296. )
  297. tf32_timer = benchmark.Timer(
  298. stmt="torch.mm(matrixA, matrixB)",
  299. setup="torch.backends.cuda.matmul.allow_tf32 = True",
  300. globals={"matrixA": matrixA, "matrixB": matrixB},
  301. )
  302. torch.backends.cuda.matmul.allow_tf32 = False
  303. fp32_time = fp32_timer.timeit(10).mean
  304. tf32_time = tf32_timer.timeit(10).mean
  305. shapes_factor_map[shape] = tf32_time / fp32_time
  306. return shapes_factor_map
  307. class OptimizerSingleTensorPattern(Pattern):
  308. """
  309. This pattern identifies if we are using the single-tensor version of an optimizer.
  310. example:
  311. optimizer = torch.optim.SGD(model.parameters(), lr=0.1)
  312. By adding foreach=True to enable multi-tensor optimizer, we can gain speedup when
  313. the kernels are relatively small.
  314. Pattern:
  315. XXXXX: _single_tenser_<OPTIMIZER_NAME>
  316. Algorithm:
  317. String match
  318. """
  319. def __init__(self, prof: profile, should_benchmark: bool = False) -> None:
  320. super().__init__(prof, should_benchmark)
  321. self.name = "Optimizer Single Tensor Pattern"
  322. self.optimizers_with_foreach = ["adam", "sgd", "adamw"]
  323. self.description = (
  324. "Detected optimizer running with single tensor implementation. "
  325. "Please enable multi tensor implementation by passing 'foreach=True' into optimizer."
  326. )
  327. self.url = ""
  328. def match(self, event: _ProfilerEvent) -> bool:
  329. for optimizer in self.optimizers_with_foreach:
  330. if event.name.endswith(f"_single_tensor_{optimizer}"):
  331. return True
  332. return False
  333. class SynchronizedDataLoaderPattern(Pattern):
  334. """
  335. This pattern identifies if we are using num_workers=0 in DataLoader.
  336. example:
  337. torch.utils.data.DataLoader(dataset, batch_size=batch_size)
  338. Add num_workers=N to the arguments. N depends on system configuration.
  339. Pattern:
  340. dataloader.py(...): __iter__
  341. dataloader.py(...): _get_iterator
  342. NOT dataloader.py(...): check_worker_number_rationality
  343. Algorithm:
  344. If we don't see check_worker_number_rationality call in the dataloader __iter__,
  345. It is not an asynchronous dataloader.
  346. """
  347. def __init__(self, prof: profile, should_benchmark: bool = False) -> None:
  348. super().__init__(prof, should_benchmark)
  349. self.name = "Synchronized DataLoader Pattern"
  350. self.description = (
  351. "Detected DataLoader running with synchronized implementation. "
  352. "Please enable asynchronous dataloading by setting num_workers > 0 when initializing DataLoader."
  353. )
  354. self.url = (
  355. "https://pytorch.org/tutorials/recipes/recipes/tuning_guide.html"
  356. "#enable-async-data-loading-and-augmentation"
  357. )
  358. def match(self, event: _ProfilerEvent) -> bool:
  359. def is_dataloader_function(name: str, function_name: str):
  360. return name.startswith(
  361. os.path.join("torch", "utils", "data", "dataloader.py")
  362. ) and name.endswith(function_name)
  363. # TODO: fixme! Due to lifetime issues of the function name, this field might
  364. # actually point to an already freed string when the even is a PyCall.
  365. # Just silently skip this to unblock testing.
  366. try:
  367. event.name
  368. except UnicodeDecodeError:
  369. return False
  370. if not is_dataloader_function(event.name, "__iter__"):
  371. return False
  372. if not event.children:
  373. return False
  374. event = event.children[0]
  375. if not is_dataloader_function(event.name, "_get_iterator"):
  376. return False
  377. if not event.children:
  378. return False
  379. event = event.children[0]
  380. return not is_dataloader_function(event.name, "check_worker_number_rationality")
  381. # TODO: We should also check if the loader is bottleneck.
  382. class GradNotSetToNonePattern(Pattern):
  383. """
  384. This pattern identifies if we are not setting grad to None in zero_grad.
  385. example:
  386. optimizer.zero_grad()
  387. By setting set_to_none=True, we can gain speedup
  388. Pattern:
  389. XXXXX: _zero_grad
  390. NOT aten::zeros
  391. aten::zero_
  392. aten::zero_ is called on each parameter in the model.
  393. We also want to make sure it is not called by aten::zeros.
  394. Algorithm:
  395. String match
  396. """
  397. def __init__(self, prof: profile, should_benchmark: bool = False) -> None:
  398. super().__init__(prof, should_benchmark)
  399. self.name = "Gradient Set To Zero Instead of None Pattern"
  400. self.description = (
  401. "Detected gradient set to zero instead of None. "
  402. "Please add 'set_to_none=True' when calling zero_grad()."
  403. )
  404. self.url = (
  405. "https://pytorch.org/tutorials/recipes/recipes/tuning_guide.html"
  406. "#disable-gradient-calculation-for-validation-or-inference"
  407. )
  408. def match(self, event: _ProfilerEvent) -> bool:
  409. if not event.name.endswith(": zero_grad"):
  410. return False
  411. if not event.children:
  412. return False
  413. for sub_event in traverse_dfs(event.children):
  414. if (
  415. sub_event.name == "aten::zero_"
  416. and sub_event.parent.name != "aten::zeros"
  417. ):
  418. return True
  419. # TODO: We should also check if the optimizer's numerical behavior will change.
  420. return False
  421. class Conv2dBiasFollowedByBatchNorm2dPattern(Pattern):
  422. """
  423. This pattern identifies if we are enabling bias in Conv2d which is followed by BatchNorm2d.
  424. Bias doesn't do anything when followed by batchnorm.
  425. Pattern:
  426. nn.Module: Conv2d | nn.Module: BatchNorm2d
  427. ...
  428. aten::conv2d AND dtype of third argument is not null
  429. The third argument is the bias
  430. Algorithm:
  431. String match
  432. """
  433. def __init__(self, prof: profile, should_benchmark: bool = False) -> None:
  434. super().__init__(prof, should_benchmark)
  435. self.name = "Enabling Bias in Conv2d Followed By BatchNorm Pattern"
  436. self.description = "Detected bias enabled in Conv2d that is followed by BatchNorm2d. Please set 'bias=False' in Conv2d."
  437. self.url = (
  438. "https://pytorch.org/tutorials/recipes/recipes/tuning_guide.html"
  439. "#disable-bias-for-convolutions-directly-followed-by-a-batch-norm"
  440. )
  441. @property
  442. def skip(self):
  443. return self.prof.record_shapes is False or super().skip
  444. def match(self, event: _ProfilerEvent):
  445. if event.name != "aten::conv2d":
  446. return False
  447. if len(input_dtypes(event)) < 3 or input_dtypes(event)[2] is None:
  448. return False
  449. # This means bias=True
  450. event = self.go_up_until(
  451. event, lambda e: e.name.startswith("nn.Module: Conv2d")
  452. )
  453. if not event:
  454. return False
  455. event = self.next_of(event)
  456. if not event:
  457. return False
  458. return event.name.startswith("nn.Module: BatchNorm2d")
  459. class MatMulDimInFP16Pattern(Pattern):
  460. def __init__(self, prof: profile, should_benchmark: bool = False) -> None:
  461. super().__init__(prof, should_benchmark)
  462. self.name = "Matrix Multiplication Dimension Not Aligned Pattern"
  463. self.description = "Detected matmul with dimension not aligned. Please use matmul with aligned dimension."
  464. self.url = "https://pytorch.org/tutorials/recipes/recipes/tuning_guide.html#use-mixed-precision-and-amp"
  465. @property
  466. def skip(self) -> bool:
  467. return not self.prof.with_stack or not self.prof.record_shapes
  468. def match(self, event: _ProfilerEvent) -> bool:
  469. def mutiple_of(shapes, multiple):
  470. return all(dim % multiple == 0 for shape in shapes for dim in shape[-2:])
  471. if event.name not in ("aten::mm", "aten::bmm", "aten::addmm"):
  472. return False
  473. if not input_dtypes(event):
  474. return False
  475. arg_dtype = input_dtypes(event)[0]
  476. if arg_dtype in (torch.bfloat16, torch.half) and not mutiple_of(
  477. input_shapes(event), 8
  478. ):
  479. return True
  480. return False
  481. def benchmark(self, events: list[_ProfilerEvent]):
  482. def closest_multiple(shapes, multiple):
  483. return [multiple * math.ceil(shape / multiple) for shape in shapes]
  484. shapes_factor_map = {input_shapes(event): 0.0 for event in events}
  485. for shape in shapes_factor_map:
  486. matrixA = torch.randn(shape[0], device="cuda", dtype=torch.float16)
  487. matrixB = torch.randn(shape[1], device="cuda", dtype=torch.float16)
  488. not_aligned_dim_timer = benchmark.Timer(
  489. stmt="torch.mm(matrixA, matrixB)",
  490. globals={"matrixA": matrixA, "matrixB": matrixB},
  491. )
  492. matrixA = torch.randn(
  493. closest_multiple(shape[0], 8), device="cuda", dtype=torch.float16
  494. )
  495. matrixB = torch.randn(
  496. closest_multiple(shape[1], 8), device="cuda", dtype=torch.float16
  497. )
  498. aligned_dim_timer = benchmark.Timer(
  499. stmt="torch.mm(matrixA, matrixB)",
  500. globals={"matrixA": matrixA, "matrixB": matrixB},
  501. )
  502. not_aligned_dim_time = not_aligned_dim_timer.timeit(10).mean
  503. aligned_dim_time = aligned_dim_timer.timeit(10).mean
  504. shapes_factor_map[shape] = aligned_dim_time / not_aligned_dim_time
  505. return shapes_factor_map
  506. def source_code_location(event: Optional[_ProfilerEvent]) -> str:
  507. while event:
  508. if event.tag == _EventType.PyCall or event.tag == _EventType.PyCCall:
  509. assert isinstance(
  510. event.extra_fields, (_ExtraFields_PyCall, _ExtraFields_PyCCall)
  511. )
  512. if not event.extra_fields.caller.file_name.startswith("torch" + os.sep):
  513. return f"{event.extra_fields.caller.file_name}:{event.extra_fields.caller.line_number}"
  514. event = event.parent
  515. return "No source code location found"
  516. def input_shapes(event: _ProfilerEvent):
  517. assert isinstance(event.extra_fields, _ExtraFields_TorchOp)
  518. return tuple(tuple(getattr(i, "sizes", ())) for i in event.extra_fields.inputs)
  519. def input_dtypes(event: _ProfilerEvent):
  520. assert isinstance(event.extra_fields, _ExtraFields_TorchOp)
  521. return tuple(getattr(i, "dtype", None) for i in event.extra_fields.inputs)
  522. def report_all_anti_patterns(
  523. prof,
  524. should_benchmark: bool = False,
  525. print_enable: bool = True,
  526. json_report_dir: Optional[str] = None,
  527. ) -> None:
  528. report_dict: dict = {}
  529. anti_patterns = [
  530. ExtraCUDACopyPattern(prof, should_benchmark),
  531. # ForLoopIndexingPattern(prof, should_benchmark),
  532. FP32MatMulPattern(prof, should_benchmark),
  533. OptimizerSingleTensorPattern(prof, should_benchmark),
  534. SynchronizedDataLoaderPattern(prof, should_benchmark),
  535. GradNotSetToNonePattern(prof, should_benchmark),
  536. Conv2dBiasFollowedByBatchNorm2dPattern(prof, should_benchmark),
  537. MatMulDimInFP16Pattern(prof, should_benchmark),
  538. ]
  539. reported = set()
  540. summaries = []
  541. message_list = [f"{'-' * 40}TorchTidy Report{'-' * 40}"]
  542. message_list.append("Matched Events:")
  543. for anti_pattern in anti_patterns:
  544. matched_events = anti_pattern.matched_events()
  545. if not matched_events:
  546. continue
  547. summaries.append(anti_pattern.summary(matched_events))
  548. for event in matched_events:
  549. report_msg = anti_pattern.report(event)
  550. if report_msg not in reported:
  551. message_list.append(report_msg)
  552. reported.add(report_msg)
  553. src_location, line_no = source_code_location(event).split(":")
  554. report_dict.setdefault(src_location, []).append(
  555. {
  556. "line_number": int(line_no),
  557. "name": anti_pattern.name,
  558. "url": anti_pattern.url,
  559. "message": anti_pattern.description,
  560. }
  561. )
  562. if json_report_dir is not None:
  563. json_report_path = os.path.join(json_report_dir, "torchtidy_report.json")
  564. if os.path.exists(json_report_path):
  565. with open(json_report_path) as f:
  566. exisiting_report = json.load(f)
  567. exisiting_report.update(report_dict)
  568. report_dict = exisiting_report
  569. with open(json_report_path, "w") as f:
  570. json.dump(report_dict, f, indent=4)
  571. message_list.append("Summary:")
  572. message_list += summaries
  573. message_list.append(f"{'-' * 40}TorchTidy Report{'-' * 40}")
  574. if print_enable:
  575. print("\n".join(message_list))