|
1 | 1 | # Copyright (c) OpenMMLab. All rights reserved.
|
2 | 2 | import pytest
|
3 | 3 | import torch
|
| 4 | +import numpy as np |
4 | 5 |
|
5 | 6 | from mmcv.ops import three_interpolate
|
6 | 7 | from mmcv.utils import IS_CUDA_AVAILABLE, IS_NPU_AVAILABLE
|
@@ -95,3 +96,89 @@ def test_three_interpolate(dtype, device):
|
95 | 96 | device=device)
|
96 | 97 |
|
97 | 98 | assert torch.allclose(output, expected_output, 1e-3, 1e-4)
|
| 99 | + |
| 100 | + |
| 101 | +def three_interpolate_forward_gloden(features, idx, weight): |
| 102 | + bs, cs, ms = features.shape |
| 103 | + ns = idx.shape[1] |
| 104 | + |
| 105 | + dtype = features.dtype |
| 106 | + if dtype == np.float16: |
| 107 | + features = features.astype(np.float32) |
| 108 | + weight = weight.astype(np.float32) |
| 109 | + |
| 110 | + output = np.zeros((bs, cs, ns), dtype=np.float) |
| 111 | + for b in range(bs): |
| 112 | + for c in range(cs): |
| 113 | + for n in range(ns): |
| 114 | + output[b][c][n] = \ |
| 115 | + features[b][c][idx[b][n][0]] * weight[b][n][0] \ |
| 116 | + + features[b][c][idx[b][n][1]] * weight[b][n][1] \ |
| 117 | + + features[b][c][idx[b][n][2]] * weight[b][n][2] |
| 118 | + return output |
| 119 | + |
| 120 | + |
| 121 | +def three_interpolate_backward_gloden(grad_output, idx, weight, features): |
| 122 | + bs, cs, ns = grad_output.shape |
| 123 | + ms = features.shape[2] |
| 124 | + |
| 125 | + dtype = features.dtype |
| 126 | + if dtype == np.float16: |
| 127 | + features = features.astype(np.float32) |
| 128 | + weight = weight.astype(np.float32) |
| 129 | + |
| 130 | + grad_point = np.zeros((bs, cs, ms), dtype=np.float) |
| 131 | + for b in range(bs): |
| 132 | + for c in range(cs): |
| 133 | + for n in range(ns): |
| 134 | + grad_point[b][c][idx[b][n][0]] = \ |
| 135 | + grad_point[b][c][idx[b][n][0]] + \ |
| 136 | + grad_output[b][c][n] * weight[b][n][0] |
| 137 | + grad_point[b][c][idx[b][n][1]] = \ |
| 138 | + grad_point[b][c][idx[b][n][1]] + \ |
| 139 | + grad_output[b][c][n] * weight[b][n][1] |
| 140 | + grad_point[b][c][idx[b][n][2]] = \ |
| 141 | + grad_point[b][c][idx[b][n][2]] + \ |
| 142 | + grad_output[b][c][n] * weight[b][n][2] |
| 143 | + return grad_point |
| 144 | + |
| 145 | + |
| 146 | +def torch_type_trans(dtype): |
| 147 | + if dtype == torch.half: |
| 148 | + return np.float16 |
| 149 | + elif dtype == torch.float: |
| 150 | + return np.float32 |
| 151 | + else: |
| 152 | + return np.float64 |
| 153 | + |
| 154 | + |
| 155 | +@pytest.mark.parametrize('dtype', [torch.half, torch.float]) |
| 156 | +@pytest.mark.parametrize('device', [ |
| 157 | + pytest.param( |
| 158 | + 'npu', |
| 159 | + marks=pytest.mark.skipif( |
| 160 | + not IS_NPU_AVAILABLE, reason='requires NPU support')) |
| 161 | +]) |
| 162 | +@pytest.mark.parametrize('shape', [(2, 5, 6, 6), (10, 10, 10, 10), |
| 163 | + (20, 21, 13, 4), (2, 10, 2, 18), |
| 164 | + (10, 602, 910, 200), (600, 100, 300, 101)]) |
| 165 | +def test_three_interpolate_npu_dynamic_shape(dtype, device, shape): |
| 166 | + bs = shape[0] |
| 167 | + cs = shape[1] |
| 168 | + ms = shape[2] |
| 169 | + ns = shape[3] |
| 170 | + |
| 171 | + features = np.random.uniform(-10.0, 10.0, |
| 172 | + (bs, cs, ms)).astype(torch_type_trans(dtype)) |
| 173 | + idx = np.random.uniform(0, ms, size=(bs, ns, 3), dtype=np.int32) |
| 174 | + weight = np.random.uniform(-10.0, |
| 175 | + 10.0 (bs, ns, |
| 176 | + 3)).astype(torch_type_trans(dtype)) |
| 177 | + |
| 178 | + features_npu = torch.tensor(features, dtype=dtype).to(device) |
| 179 | + idx_npu = torch.tensor(idx, dtype=torch.int32).to(device) |
| 180 | + weight_npu = torch.tensor(weight, dtype=dtype).to(device) |
| 181 | + |
| 182 | + expected_output = three_interpolate_forward_gloden(features, idx, weight) |
| 183 | + output = three_interpolate(features_npu, idx_npu, weight_npu) |
| 184 | + assert np.allclose(output.cpu().numpy(), expected_output, 1e-3, 1e-4) |
0 commit comments