Skip to content

Commit 92c8880

Browse files
authored
Merge pull request #1124 from calad0i/general_transpose
Add general transpose for vivado/vitis
2 parents fd594e0 + bbc316e commit 92c8880

File tree

8 files changed

+215
-102
lines changed

8 files changed

+215
-102
lines changed

hls4ml/backends/vivado/passes/reshaping_templates.py

Lines changed: 67 additions & 15 deletions
Original file line numberDiff line numberDiff line change
@@ -1,3 +1,7 @@
1+
from math import prod
2+
3+
import numpy as np
4+
15
from hls4ml.backends.template import FunctionCallTemplate, LayerConfigTemplate
26
from hls4ml.model.layers import Resize, Transpose, ZeroPadding1D, ZeroPadding2D
37

@@ -97,16 +101,64 @@ def format(self, node):
97101

98102
# Transpose templates
99103

100-
transpose_config_template = """struct config{index} : nnet::transpose_config {{
101-
static const unsigned depth = {depth};
102-
static const unsigned height = {height};
103-
static const unsigned width = {width};
104-
static constexpr unsigned perm[3] = {{{perm_str}}};
105-
}};\n"""
106104

107-
transpose_function_template = 'nnet::transpose_{dim}<{input_t}, {output_t}, {config}>({input}, {output});'
108-
109-
transpose_include_list = ['nnet_utils/nnet_array.h', 'nnet_utils/nnet_stream.h']
105+
transpose_include_list = ['nnet_utils/nnet_transpose.h', 'nnet_utils/nnet_transpose_stream.h']
106+
107+
transpose_config_template = """struct {config_name} {{
108+
static const unsigned dims = {dims};
109+
static const unsigned N = {N};
110+
static const unsigned* const from_shape;
111+
static const unsigned* const to_shape;
112+
static const unsigned* const perm;
113+
static const unsigned* const perm_strides;
114+
}};
115+
116+
unsigned {config_name}_from_shape[{dims}] = {{{from_shape}}};
117+
unsigned {config_name}_to_shape[{dims}] = {{{to_shape}}};
118+
unsigned {config_name}_perm[{dims}] = {{{perm}}};
119+
unsigned {config_name}_perm_strides[{dims}] = {{{perm_strides}}};
120+
121+
const unsigned* const {config_name}::from_shape = {config_name}_from_shape;
122+
const unsigned* const {config_name}::to_shape = {config_name}_to_shape;
123+
const unsigned* const {config_name}::perm = {config_name}_perm;
124+
const unsigned* const {config_name}::perm_strides = {config_name}_perm_strides;
125+
"""
126+
127+
transpose_function_template = 'nnet::transpose<{input_t}, {output_t}, {config_name}>({input}, {output});'
128+
129+
130+
def permute_config_gen(name: str, shape: tuple[int, ...], perm: tuple[int, ...]):
131+
"""
132+
Generate a configuration string for a permute operation. Operates by mapping the output index to input input index by:
133+
- unravel the output index
134+
- map each dimension to the corresponding stride in the input tensor, sum
135+
The operation can be expressed as:
136+
137+
new_shape = tuple(shape[i] for i in perm)
138+
strides = np.cumprod((shapes[1:] + (1,))[::-1])[::-1]
139+
perm_strides = [strides[i] for i in perm]
140+
out[index] = inp[np.dot(np.unravel_index(index, new_shape), perm_strides)]
141+
142+
Args:
143+
name (str): The name of the configuration.
144+
shape (tuple[int, ...]): The shape of the input tensor.
145+
perm (tuple[int, ...]): The permutation of the dimensions.
146+
147+
Returns:
148+
str: The formatted configuration string for the permute operation.
149+
"""
150+
new_shape = tuple(shape[i] for i in perm)
151+
strides = np.cumprod((shape[1:] + (1,))[::-1])[::-1]
152+
perm_strides = tuple(int(strides[i]) for i in perm)
153+
return transpose_config_template.format(
154+
dims=len(shape),
155+
N=prod(shape),
156+
from_shape=', '.join(str(x) for x in shape),
157+
perm=', '.join(str(x) for x in perm),
158+
perm_strides=', '.join(str(x) for x in perm_strides),
159+
to_shape=', '.join(str(x) for x in new_shape),
160+
config_name=name,
161+
)
110162

111163

112164
class TransposeConfigTemplate(LayerConfigTemplate):
@@ -115,18 +167,18 @@ def __init__(self):
115167
self.template = transpose_config_template
116168

117169
def format(self, node):
118-
params = self._default_config_params(node)
119-
120-
return self.template.format(**params)
170+
shape = tuple(node.get_input_variable().shape)
171+
perm = tuple(node.get_attr('perm'))
172+
name = f'config{node.index}'
173+
return permute_config_gen(name, shape, perm)
121174

122175

123176
class TransposeFunctionTemplate(FunctionCallTemplate):
124177
def __init__(self):
125-
super().__init__(Transpose, include_header=transpose_include_list)
126178
self.template = transpose_function_template
179+
super().__init__(Transpose, include_header=transpose_include_list)
127180

128181
def format(self, node):
129182
params = self._default_function_params(node)
130-
params['dim'] = node.get_attr('dim')
131-
183+
params['config_name'] = f'config{node.index}'
132184
return self.template.format(**params)

hls4ml/model/layers.py

Lines changed: 5 additions & 3 deletions
Original file line numberDiff line numberDiff line change
@@ -1221,8 +1221,7 @@ def initialize(self):
12211221
perm = self.get_attr('perm')
12221222
self.set_attr('dim', f'{len(inp.shape)}d')
12231223

1224-
if len(perm) > 3:
1225-
raise Exception('ERROR: Transpose of tensors with rank > 3 is not yet supported.')
1224+
# TODO: dim>3 is only supported for vivado/vitis backend
12261225

12271226
# ONNX double transpose specific, sometimes ONNX injects
12281227
# useless double transpose layers when converting
@@ -1242,11 +1241,14 @@ def initialize(self):
12421241
self.set_attr('depth', 1)
12431242
self.set_attr('height', inp.shape[0])
12441243
self.set_attr('width', inp.shape[1])
1245-
elif len(shape) > 2:
1244+
elif len(shape) == 3:
12461245
dims = [f'OUT_DEPTH_{self.index}', f'OUT_HEIGHT_{self.index}', f'OUT_WIDTH_{self.index}']
12471246
self.set_attr('depth', inp.shape[0])
12481247
self.set_attr('height', inp.shape[1])
12491248
self.set_attr('width', inp.shape[2])
1249+
elif len(shape) > 3:
1250+
# Differentiate between 2/3/3+ dim does not really appear to be needed. To be removed?
1251+
dims = [f'OUT_DIM_{i}_{self.index}' for i in range(1, len(shape) + 1)]
12501252
self.add_output_variable(shape, dims, precision=inp.type.precision)
12511253

12521254

hls4ml/model/optimizer/passes/convert_to_channels_last.py

Lines changed: 3 additions & 9 deletions
Original file line numberDiff line numberDiff line change
@@ -12,14 +12,13 @@ class ChannelsLastConverter(OptimizerPass):
1212
and adding a transpose layer for the inputs and outputs, if necessary'''
1313

1414
def match(self, node):
15+
# If this parameter has not been set, this model does not need to be converted
16+
if 'ChannelsLastConversion' not in node.model.config.config['HLSConfig']['Model']:
17+
return False # No littering of unused property
1518
if not hasattr(node, 'channels_last_converted'):
1619
return True
1720

1821
def transform(self, model, node):
19-
# If this parameter has not been set, this model does not need to be converted
20-
if 'ChannelsLastConversion' not in model.config.config['HLSConfig']['Model']:
21-
node.channels_last_converted = True
22-
return False
2322
outshape = node.get_output_variable().shape
2423

2524
if isinstance(node, Input):
@@ -103,11 +102,6 @@ def transform(self, model, node):
103102
input = previous_node.name
104103
outshape = previous_node.get_output_variable().shape
105104

106-
if (model.config.config['IOType'] == 'io_stream') and len(outshape) == 3:
107-
raise Exception(
108-
'No 3D transpose available in io_stream, this model cannot be converted to channels-last'
109-
)
110-
111105
if len(outshape) == 2:
112106
attributes = {'perm': [1, 0]}
113107
else:

hls4ml/templates/vivado/nnet_utils/nnet_array.h

Lines changed: 0 additions & 52 deletions
This file was deleted.

hls4ml/templates/vivado/nnet_utils/nnet_stream.h

Lines changed: 0 additions & 23 deletions
Original file line numberDiff line numberDiff line change
@@ -179,29 +179,6 @@ void broadcast_stream(hls::stream<data_T> &data, hls::stream<res_T> &res) {
179179
}
180180
}
181181

182-
template <class data_T, class res_T, typename CONFIG_T>
183-
void transpose_2d(hls::stream<data_T> &data, hls::stream<res_T> &res) {
184-
typename data_T::value_type data_array[CONFIG_T::height * CONFIG_T::width];
185-
#pragma HLS ARRAY_PARTITION variable=data_array complete
186-
187-
for (int i = 0; i < CONFIG_T::height * CONFIG_T::width / data_T::size; i++) {
188-
#pragma HLS PIPELINE
189-
data_T in_data = data.read();
190-
for (int j = 0; j < data_T::size; j++) {
191-
data_array[i * data_T::size + j] = typename data_T::value_type(in_data[j]);
192-
}
193-
}
194-
195-
for (int i = 0; i < CONFIG_T::height * CONFIG_T::width / res_T::size; i++) {
196-
#pragma HLS PIPELINE
197-
res_T out_data;
198-
PRAGMA_DATA_PACK(out_data)
199-
for (int j = 0; j < res_T::size; j++) {
200-
out_data[j] = typename res_T::value_type(data_array[j * data_T::size + i]);
201-
}
202-
res.write(out_data);
203-
}
204-
}
205182
} // namespace nnet
206183

207184
#endif
Lines changed: 39 additions & 0 deletions
Original file line numberDiff line numberDiff line change
@@ -0,0 +1,39 @@
1+
#ifndef NNET_PERMUTE_H_
2+
#define NNET_PERMUTE_H_
3+
4+
namespace nnet {
5+
6+
struct transpose_config {
7+
static const unsigned dims;
8+
static const unsigned N;
9+
// vivado/vitis hls can't index constexpr array for some reason
10+
// and vivado hls don't like template recursion either (vitis is fine)
11+
// thus this appears to be the only workaround (or overkill it with codegen)
12+
static const unsigned *const from_shape;
13+
static const unsigned *const to_shape;
14+
static const unsigned *const perm;
15+
static const unsigned *const perm_strides;
16+
};
17+
18+
template <typename CONFIG_T> unsigned transfer_idx(int index) {
19+
// Given output idx in c-order flat array, return input idx
20+
int idx = 0;
21+
for (int i = CONFIG_T::dims - 1; i >= 0; i--) {
22+
idx += (index % CONFIG_T::to_shape[i]) * CONFIG_T::perm_strides[i];
23+
index /= CONFIG_T::to_shape[i];
24+
}
25+
return idx;
26+
}
27+
28+
template <typename data_T, typename res_T, typename CONFIG_T>
29+
void transpose(const data_T data[CONFIG_T::N], res_T res[CONFIG_T::N]) {
30+
for (int i = 0; i < CONFIG_T::N; i++) {
31+
#pragma HLS UNROLL
32+
int idx = transfer_idx<CONFIG_T>(i);
33+
res[i] = data[idx];
34+
}
35+
}
36+
37+
} // namespace nnet
38+
39+
#endif
Lines changed: 67 additions & 0 deletions
Original file line numberDiff line numberDiff line change
@@ -0,0 +1,67 @@
1+
#ifndef NNET_TRANSPOSE_STREAM_H
2+
#define NNET_TRANSPOSE_STREAM_H
3+
4+
#include "hls_stream.h"
5+
#include "nnet_transpose.h"
6+
#include <type_traits>
7+
8+
namespace nnet {
9+
10+
template <typename data_T, typename res_T, typename CONFIG_T>
11+
typename std::enable_if<CONFIG_T::dims == 2, void>::type transpose(hls::stream<data_T> &data, hls::stream<res_T> &res) {
12+
// #pragma HLS INLINE RECURSIVE
13+
typename data_T::value_type data_array[CONFIG_T::N];
14+
#pragma HLS ARRAY_PARTITION variable=data_array complete
15+
16+
for (int i = 0; i < CONFIG_T::N / data_T::size; i++) {
17+
#pragma HLS PIPELINE
18+
data_T in_data = data.read();
19+
for (int j = 0; j < data_T::size; j++) {
20+
#pragma HLS UNROLL
21+
data_array[i * data_T::size + j] = typename data_T::value_type(in_data[j]);
22+
}
23+
}
24+
25+
for (int i = 0; i < CONFIG_T::N / res_T::size; i++) {
26+
#pragma HLS PIPELINE
27+
res_T out_data;
28+
PRAGMA_DATA_PACK(out_data)
29+
for (int j = 0; j < res_T::size; j++) {
30+
#pragma HLS UNROLL
31+
out_data[j] = typename res_T::value_type(data_array[j * CONFIG_T::from_shape[1] + i]);
32+
}
33+
res.write(out_data);
34+
}
35+
}
36+
37+
// This sfinae is for vivado_hls, which has some overhead using the transfer_idx in io_stream.
38+
// In vitis both performs exactly the same, thus this is not removed out of convenience.
39+
template <typename data_T, typename res_T, typename CONFIG_T>
40+
typename std::enable_if<CONFIG_T::dims != 2, void>::type transpose(hls::stream<data_T> &data, hls::stream<res_T> &res) {
41+
// #pragma HLS INLINE RECURSIVE
42+
typename data_T::value_type data_array[CONFIG_T::N];
43+
#pragma HLS ARRAY_PARTITION variable=data_array complete
44+
45+
for (int i = 0; i < CONFIG_T::N / data_T::size; i++) {
46+
#pragma HLS PIPELINE
47+
data_T in_data = data.read();
48+
for (int j = 0; j < data_T::size; j++) {
49+
#pragma HLS UNROLL
50+
data_array[i * data_T::size + j] = typename data_T::value_type(in_data[j]);
51+
}
52+
}
53+
54+
for (int i = 0; i < CONFIG_T::N / res_T::size; i++) {
55+
#pragma HLS PIPELINE
56+
res_T out_data;
57+
PRAGMA_DATA_PACK(out_data)
58+
for (int j = 0; j < res_T::size; j++) {
59+
#pragma HLS UNROLL
60+
out_data[j] = typename res_T::value_type(data_array[transfer_idx<CONFIG_T>(i * res_T::size + j)]);
61+
}
62+
res.write(out_data);
63+
}
64+
}
65+
66+
} // namespace nnet
67+
#endif

test/pytest/test_transpose_concat.py

Lines changed: 34 additions & 0 deletions
Original file line numberDiff line numberDiff line change
@@ -54,3 +54,37 @@ def test_accuracy(data, keras_model, hls_model):
5454
y_hls4ml = hls_model.predict(X).reshape(y_keras.shape)
5555
# "accuracy" of hls4ml predictions vs keras
5656
np.testing.assert_allclose(y_keras, y_hls4ml, rtol=0, atol=1e-04, verbose=True)
57+
58+
59+
@pytest.fixture(scope='module')
60+
def keras_model_highdim():
61+
inp = Input(shape=(2, 3, 4, 5, 6), name='input_1')
62+
out = Permute((3, 5, 4, 1, 2))(inp)
63+
model = Model(inputs=inp, outputs=out)
64+
return model
65+
66+
67+
@pytest.fixture(scope='module')
68+
def data_highdim():
69+
X = np.random.randint(-128, 127, (100, 2, 3, 4, 5, 6)) / 128
70+
X = X.astype(np.float32)
71+
return X
72+
73+
74+
@pytest.mark.parametrize('io_type', ['io_stream', 'io_parallel'])
75+
@pytest.mark.parametrize('backend', ['Vivado', 'Vitis'])
76+
def test_highdim_permute(data_highdim, keras_model_highdim, io_type, backend):
77+
X = data_highdim
78+
model = keras_model_highdim
79+
80+
model_hls = hls4ml.converters.convert_from_keras_model(
81+
model,
82+
io_type=io_type,
83+
backend=backend,
84+
output_dir=str(test_root_path / f'hls4mlprj_highdim_transpose_{backend}_{io_type}'),
85+
)
86+
model_hls.compile()
87+
y_keras = model.predict(X)
88+
y_hls4ml = model_hls.predict(X).reshape(y_keras.shape) # type: ignore
89+
90+
assert np.all(y_keras == y_hls4ml)

0 commit comments

Comments
 (0)