Skip to content

Commit 97c187d

Browse files
authored
Merge branch 'main' into oneapi_backend/experiment
2 parents 7e028e6 + 2366f9b commit 97c187d

File tree

9 files changed

+138
-136
lines changed

9 files changed

+138
-136
lines changed

CITATION.cff

Lines changed: 2 additions & 1 deletion
Original file line numberDiff line numberDiff line change
@@ -4,7 +4,8 @@ type: software
44
authors:
55
- given-names: "FastML Team"
66
title: "hls4ml"
7-
version: "v1.0.0"
7+
version: "v1.1.0"
8+
date-released: "2025-03-17"
89
doi: 10.5281/zenodo.1201549
910
repository-code: "https://github.com/fastmachinelearning/hls4ml"
1011
url: "https://fastmachinelearning.org/hls4ml"

README.md

Lines changed: 2 additions & 2 deletions
Original file line numberDiff line numberDiff line change
@@ -73,9 +73,9 @@ If you use this software in a publication, please cite the software
7373
@software{fastml_hls4ml,
7474
author = {{FastML Team}},
7575
title = {fastmachinelearning/hls4ml},
76-
year = 2024,
76+
year = 2025,
7777
publisher = {Zenodo},
78-
version = {v1.0.0},
78+
version = {v1.1.0},
7979
doi = {10.5281/zenodo.1201549},
8080
url = {https://github.com/fastmachinelearning/hls4ml}
8181
}

docs/advanced/extension.rst

Lines changed: 3 additions & 0 deletions
Original file line numberDiff line numberDiff line change
@@ -18,6 +18,9 @@ To implement a custom layer in ``hls4ml`` with the extension API, the required c
1818
* Function config template
1919
* Registration of layer, source code, and templates
2020

21+
.. note::
22+
currently, then extension API supports keras models. Support for pytorch models is in development.
23+
2124
Complete example
2225
================
2326

docs/intro/setup.rst

Lines changed: 1 addition & 7 deletions
Original file line numberDiff line numberDiff line change
@@ -20,14 +20,8 @@ If you want to use our :doc:`profiling <../advanced/profiling>` toolbox, you mig
2020
2121
pip install hls4ml[profiling]
2222
23-
``hls4ml`` is also available as a ``conda`` package in the ``conda-forge`` repository. To install, run:
24-
2523
.. warning::
26-
Version of hls4ml available on ``conda-forge`` is outdated, we recommend installing with ``pip`` to get the latest version.
27-
28-
.. code-block::
29-
30-
conda install -c conda-forge hls4ml
24+
Previously, versions of hls4ml were made available on ``conda-forge``. These are outdated and should NOT be used. Installing with ``pip`` is currently the only supported method.
3125

3226
Development version
3327
-------------------

docs/intro/status.rst

Lines changed: 5 additions & 4 deletions
Original file line numberDiff line numberDiff line change
@@ -89,14 +89,15 @@ A summary of the on-going status of the ``hls4ml`` tool is in the table below.
8989

9090
Other feature notes:
9191

92-
* ``hls4ml`` is tested on Linux, and supports
92+
* ``hls4ml`` is tested on the following platforms. Newer versions might work just fine, but try at your own risk.
9393
* Vivado HLS versions 2018.2 to 2020.1
94-
* Intel HLS versions 20.1 to 21.4
95-
* Vitis HLS versions 2022.2 to 2024.1
94+
* Intel HLS versions 20.1 to 21.4, versions \> 21.4 have not been tested.
95+
* Vitis HLS versions 2022.2 to 2024.1. Versions \<= 2022.1 are known not to work.
9696
* Catapult HLS versions 2024.1_1 to 2024.2
9797
* oneAPI versions 2024.1 to 2025.0
9898

99-
* Windows and macOS are not supported
99+
* ``hls4ml`` supports Linux and requires python \>=3.10. hlsml does not require a specific Linux distribution version and we recommended to follow the requirements of the HLS tool you are using.
100+
* Windows and macOS are not supported. Setting up ``hls4ml`` on these platforms, for example using the Windows Subsystem for Linux (WSL) should be possible, but we do not provide support for such use cases.
100101
* BDT support has moved to the `Conifer <https://github.com/thesps/conifer>`__ package
101102

102103
Example Models

hls4ml/model/graph.py

Lines changed: 4 additions & 7 deletions
Original file line numberDiff line numberDiff line change
@@ -542,14 +542,14 @@ def remove_node(self, node, rewire=True):
542542
if len(inputs) > 1 or len(outputs) > 1:
543543
raise Exception('Cannot delete a node with multiple inputs/outputs')
544544

545-
if len(inputs) == 1:
545+
if len(outputs) == 1 and len(inputs) == 1:
546+
546547
# Connect inputs -> $outputs
547-
if node.name in self.outputs:
548+
if node.outputs[0] in self.outputs:
548549
msg = f'Remove leaf node {node.name} will connect its input node {inputs[0]} to output, but it already is.'
549550
assert inputs[0] not in self.outputs, msg
550-
self.outputs = [inputs[0] if name == node.name else name for name in self.outputs]
551+
self.outputs = [inputs[0] if name == node.outputs[0] else name for name in self.outputs]
551552

552-
if len(outputs) == 1 and len(inputs) == 1:
553553
inp_var = node.get_input_variable()
554554
out_var = node.get_output_variable()
555555

@@ -565,9 +565,6 @@ def remove_node(self, node, rewire=True):
565565
if outputs[0] == nxt_inp:
566566
next_node.inputs[i] = inputs[0]
567567

568-
if node.outputs[0] in self.outputs:
569-
prev_node = node.get_input_node(node.inputs[0])
570-
self.outputs[self.outputs.index(node.outputs[0])] = prev_node.outputs[0]
571568
del self.output_vars[node.outputs[0]]
572569
del self.graph[node.name]
573570

hls4ml/model/profiling.py

Lines changed: 2 additions & 111 deletions
Original file line numberDiff line numberDiff line change
@@ -393,121 +393,12 @@ def activations_keras(model, X, fmt='longform', plot='boxplot'):
393393

394394

395395
def weights_torch(model, fmt='longform', plot='boxplot'):
396+
from hls4ml.utils.profiling_utils import WeightsTorch
397+
396398
wt = WeightsTorch(model, fmt, plot)
397399
return wt.get_weights()
398400

399401

400-
def _torch_batchnorm(layer):
401-
weights = list(layer.parameters())
402-
epsilon = layer.eps
403-
404-
gamma = weights[0]
405-
beta = weights[1]
406-
if layer.track_running_stats:
407-
mean = layer.running_mean
408-
var = layer.running_var
409-
else:
410-
mean = torch.tensor(np.ones(20))
411-
var = torch.tensor(np.zeros(20))
412-
413-
scale = gamma / np.sqrt(var + epsilon)
414-
bias = beta - gamma * mean / np.sqrt(var + epsilon)
415-
416-
return [scale, bias], ['s', 'b']
417-
418-
419-
def _torch_layer(layer):
420-
return list(layer.parameters()), ['w', 'b']
421-
422-
423-
def _torch_rnn(layer):
424-
return list(layer.parameters()), ['w_ih_l0', 'w_hh_l0', 'b_ih_l0', 'b_hh_l0']
425-
426-
427-
torch_process_layer_map = defaultdict(
428-
lambda: _torch_layer,
429-
{
430-
'BatchNorm1d': _torch_batchnorm,
431-
'BatchNorm2d': _torch_batchnorm,
432-
'RNN': _torch_rnn,
433-
'LSTM': _torch_rnn,
434-
'GRU': _torch_rnn,
435-
},
436-
)
437-
438-
439-
class WeightsTorch:
440-
def __init__(self, model: torch.nn.Module, fmt: str = 'longform', plot: str = 'boxplot') -> None:
441-
self.model = model
442-
self.fmt = fmt
443-
self.plot = plot
444-
self.registered_layers = list()
445-
self._find_layers(self.model, self.model.__class__.__name__)
446-
447-
def _find_layers(self, model, module_name):
448-
for name, module in model.named_children():
449-
if isinstance(module, (torch.nn.Sequential, torch.nn.ModuleList)):
450-
self._find_layers(module, module_name + "." + name)
451-
elif isinstance(module, (torch.nn.Module)) and self._is_parameterized(module):
452-
if len(list(module.named_children())) != 0:
453-
# custom nn.Module, continue search
454-
self._find_layers(module, module_name + "." + name)
455-
else:
456-
self._register_layer(module_name + "." + name)
457-
458-
def _is_registered(self, name: str) -> bool:
459-
return name in self.registered_layers
460-
461-
def _register_layer(self, name: str) -> None:
462-
if self._is_registered(name) is False:
463-
self.registered_layers.append(name)
464-
465-
def _is_parameterized(self, module: torch.nn.Module) -> bool:
466-
return any(p.requires_grad for p in module.parameters())
467-
468-
def _get_weights(self) -> pandas.DataFrame | list[dict]:
469-
if self.fmt == 'longform':
470-
data = {'x': [], 'layer': [], 'weight': []}
471-
elif self.fmt == 'summary':
472-
data = []
473-
for layer_name in self.registered_layers:
474-
layer = self._get_layer(layer_name, self.model)
475-
name = layer.__class__.__name__
476-
weights, suffix = torch_process_layer_map[layer.__class__.__name__](layer)
477-
for i, w in enumerate(weights):
478-
label = f'{name}/{suffix[i]}'
479-
w = weights[i].detach().numpy()
480-
w = w.flatten()
481-
w = abs(w[w != 0])
482-
n = len(w)
483-
if n == 0:
484-
print(f'Weights for {name} are only zeros, ignoring.')
485-
break
486-
if self.fmt == 'longform':
487-
data['x'].extend(w.tolist())
488-
data['layer'].extend([name] * n)
489-
data['weight'].extend([label] * n)
490-
elif self.fmt == 'summary':
491-
data.append(array_to_summary(w, fmt=self.plot))
492-
data[-1]['layer'] = name
493-
data[-1]['weight'] = label
494-
495-
if self.fmt == 'longform':
496-
data = pandas.DataFrame(data)
497-
return data
498-
499-
def get_weights(self) -> pandas.DataFrame | list[dict]:
500-
return self._get_weights()
501-
502-
def get_layers(self) -> list[str]:
503-
return self.registered_layers
504-
505-
def _get_layer(self, layer_name: str, module: torch.nn.Module) -> torch.nn.Module:
506-
for name in layer_name.split('.')[1:]:
507-
module = getattr(module, name)
508-
return module
509-
510-
511402
def activations_torch(model, X, fmt='longform', plot='boxplot'):
512403
X = torch.Tensor(X)
513404
if fmt == 'longform':

hls4ml/utils/profiling_utils.py

Lines changed: 119 additions & 0 deletions
Original file line numberDiff line numberDiff line change
@@ -0,0 +1,119 @@
1+
from collections import defaultdict
2+
3+
import numpy as np
4+
import pandas
5+
import torch
6+
7+
from hls4ml.model.profiling import array_to_summary
8+
9+
10+
def _torch_batchnorm(layer):
11+
weights = list(layer.parameters())
12+
epsilon = layer.eps
13+
14+
gamma = weights[0]
15+
beta = weights[1]
16+
if layer.track_running_stats:
17+
mean = layer.running_mean
18+
var = layer.running_var
19+
else:
20+
mean = torch.tensor(np.ones(20))
21+
var = torch.tensor(np.zeros(20))
22+
23+
scale = gamma / np.sqrt(var + epsilon)
24+
bias = beta - gamma * mean / np.sqrt(var + epsilon)
25+
26+
return [scale, bias], ['s', 'b']
27+
28+
29+
def _torch_layer(layer):
30+
return list(layer.parameters()), ['w', 'b']
31+
32+
33+
def _torch_rnn(layer):
34+
return list(layer.parameters()), ['w_ih_l0', 'w_hh_l0', 'b_ih_l0', 'b_hh_l0']
35+
36+
37+
torch_process_layer_map = defaultdict(
38+
lambda: _torch_layer,
39+
{
40+
'BatchNorm1d': _torch_batchnorm,
41+
'BatchNorm2d': _torch_batchnorm,
42+
'RNN': _torch_rnn,
43+
'LSTM': _torch_rnn,
44+
'GRU': _torch_rnn,
45+
},
46+
)
47+
48+
49+
class WeightsTorch:
50+
51+
def __init__(self, model: torch.nn.Module, fmt: str = 'longform', plot: str = 'boxplot') -> None:
52+
self.model = model
53+
self.fmt = fmt
54+
self.plot = plot
55+
self.registered_layers = list()
56+
self._find_layers(self.model, self.model.__class__.__name__)
57+
58+
def _find_layers(self, model, module_name):
59+
for name, module in model.named_children():
60+
if isinstance(module, (torch.nn.Sequential, torch.nn.ModuleList)):
61+
self._find_layers(module, module_name + "." + name)
62+
elif isinstance(module, (torch.nn.Module)) and self._is_parameterized(module):
63+
if len(list(module.named_children())) != 0:
64+
# custom nn.Module, continue search
65+
self._find_layers(module, module_name + "." + name)
66+
else:
67+
self._register_layer(module_name + "." + name)
68+
69+
def _is_registered(self, name: str) -> bool:
70+
return name in self.registered_layers
71+
72+
def _register_layer(self, name: str) -> None:
73+
if self._is_registered(name) is False:
74+
self.registered_layers.append(name)
75+
76+
def _is_parameterized(self, module: torch.nn.Module) -> bool:
77+
return any(p.requires_grad for p in module.parameters())
78+
79+
def _get_weights(self) -> pandas.DataFrame | list[dict]:
80+
if self.fmt == 'longform':
81+
data = {'x': [], 'layer': [], 'weight': []}
82+
elif self.fmt == 'summary':
83+
data = []
84+
for layer_name in self.registered_layers:
85+
layer = self._get_layer(layer_name, self.model)
86+
name = layer.__class__.__name__
87+
weights, suffix = torch_process_layer_map[layer.__class__.__name__](layer)
88+
for i, w in enumerate(weights):
89+
label = f'{name}/{suffix[i]}'
90+
w = weights[i].detach().numpy()
91+
w = w.flatten()
92+
w = abs(w[w != 0])
93+
n = len(w)
94+
if n == 0:
95+
print(f'Weights for {name} are only zeros, ignoring.')
96+
break
97+
if self.fmt == 'longform':
98+
data['x'].extend(w.tolist())
99+
data['layer'].extend([name] * n)
100+
data['weight'].extend([label] * n)
101+
elif self.fmt == 'summary':
102+
data.append(array_to_summary(w, fmt=self.plot))
103+
data[-1]['layer'] = name
104+
data[-1]['weight'] = label
105+
106+
if self.fmt == 'longform':
107+
data = pandas.DataFrame(data)
108+
return data
109+
110+
def get_weights(self) -> pandas.DataFrame | list[dict]:
111+
return self._get_weights()
112+
113+
def get_layers(self) -> list[str]:
114+
return self.registered_layers
115+
116+
def _get_layer(self, layer_name: str, module: torch.nn.Module) -> torch.nn.Module:
117+
for name in layer_name.split('.')[1:]:
118+
module = getattr(module, name)
119+
return module

test/pytest/test_qonnx.py

Lines changed: 0 additions & 4 deletions
Original file line numberDiff line numberDiff line change
@@ -402,10 +402,6 @@ def test_tiny_unet_model(tiny_unet_model, backend):
402402
@pytest.mark.parametrize('backend', ['Vitis'])
403403
@pytest.mark.parametrize('io_type', ['io_parallel', 'io_stream'])
404404
def test_simple_model(model_name, io_type, backend, request):
405-
if model_name == 'conv2d_small_mp_keras_model' and io_type == 'io_stream':
406-
# Not yet supported due to an issue with channels last conversion
407-
# There is a qonnx PR.
408-
pytest.skip()
409405
model = request.getfixturevalue(model_name)
410406
ishape = tuple(model.get_tensor_shape(model.graph.input[0].name))
411407
X = np.random.uniform(low=0, high=1, size=np.prod(ishape)).reshape(ishape)

0 commit comments

Comments
 (0)